6
6
7
7
import logging
8
8
from enum import Enum
9
- from typing import Tuple
9
+ from typing import Optional , Tuple
10
10
11
11
import torch
12
12
import torch .nn as nn
@@ -93,7 +93,7 @@ def _quantize(self, value):
93
93
)
94
94
return quantized_value , scales , zero_points
95
95
96
- def _quantize_and_update (self , input_pos , k_val , v_val ):
96
+ def _quantize_and_update (self , input_pos , k_val , v_val , indices = None ):
97
97
quantized_k_val , k_scales , k_zero_points = self ._quantize (k_val )
98
98
quantized_v_val , v_scales , v_zero_points = self ._quantize (v_val )
99
99
@@ -104,26 +104,37 @@ def _quantize_and_update(self, input_pos, k_val, v_val):
104
104
105
105
if self .use_custom_update_cache_op :
106
106
start_pos = input_pos [0 ].item ()
107
- _ = torch .ops .llama .update_cache (quantized_k_val , self .k_cache , start_pos )
108
- _ = torch .ops .llama .update_cache (k_scales , self .k_cache_scales , start_pos )
109
107
_ = torch .ops .llama .update_cache (
110
- k_zero_points , self .k_cache_zero_points , start_pos
108
+ quantized_k_val , self .k_cache , start_pos , indices
111
109
)
112
- _ = torch .ops .llama .update_cache (quantized_v_val , self .v_cache , start_pos )
113
- _ = torch .ops .llama .update_cache (v_scales , self .v_cache_scales , start_pos )
114
110
_ = torch .ops .llama .update_cache (
115
- v_zero_points , self .v_cache_zero_points , start_pos
111
+ k_scales , self .k_cache_scales , start_pos , indices
112
+ )
113
+ _ = torch .ops .llama .update_cache (
114
+ k_zero_points , self .k_cache_zero_points , start_pos , indices
115
+ )
116
+ _ = torch .ops .llama .update_cache (
117
+ quantized_v_val , self .v_cache , start_pos , indices
118
+ )
119
+ _ = torch .ops .llama .update_cache (
120
+ v_scales , self .v_cache_scales , start_pos , indices
121
+ )
122
+ _ = torch .ops .llama .update_cache (
123
+ v_zero_points , self .v_cache_zero_points , start_pos , indices
116
124
)
117
125
else :
126
+ assert indices is None , "Indices not supported for this path"
127
+ # Following is also broken because in prefill input_pos = [0]
128
+ # but we need to update some slice of cache
118
129
self .k_cache [:, input_pos ] = quantized_k_val
119
130
self .k_cache_scales [:, input_pos ] = k_scales
120
131
self .k_cache_zero_points [:, input_pos ] = k_zero_points
121
132
self .v_cache [:, input_pos ] = quantized_v_val
122
133
self .v_cache_scales [:, input_pos ] = v_scales
123
134
self .v_cache_zero_points [:, input_pos ] = v_zero_points
124
135
125
- def _update_and_return_float_values (self , input_pos , k_val , v_val ):
126
- self ._quantize_and_update (input_pos , k_val , v_val )
136
+ def _update_and_return_float_values (self , input_pos , k_val , v_val , indices = None ):
137
+ self ._quantize_and_update (input_pos , k_val , v_val , indices )
127
138
128
139
k_out = torch .ops .quantized_decomposed .dequantize_per_token (
129
140
self .k_cache ,
@@ -144,24 +155,26 @@ def _update_and_return_float_values(self, input_pos, k_val, v_val):
144
155
self .cache_fp_type ,
145
156
)
146
157
147
- # When returning float values we jsut use the last value
158
+ # When returning float values we just use the last value
148
159
# instead of dequantized value.
149
160
start_pos = input_pos [0 ].item ()
150
161
if self .use_custom_update_cache_op :
151
- _ = torch .ops .llama .update_cache (k_val , k_out , start_pos )
152
- _ = torch .ops .llama .update_cache (v_val , v_out , start_pos )
162
+ _ = torch .ops .llama .update_cache (k_val , k_out , start_pos , indices )
163
+ _ = torch .ops .llama .update_cache (v_val , v_out , start_pos , indices )
153
164
else :
154
165
k_out [:, input_pos ] = k_val
155
166
v_out [:, input_pos ] = v_val
156
167
157
168
return k_out , v_out
158
169
159
- def _update_and_return_quantized_values (self , input_pos , k_val , v_val ):
160
- self ._quantize_and_update (input_pos , k_val , v_val )
170
+ def _update_and_return_quantized_values (
171
+ self , input_pos , k_val , v_val , indices = None
172
+ ):
173
+ self ._quantize_and_update (input_pos , k_val , v_val , indices )
161
174
162
175
return self .k_cache , self .v_cache
163
176
164
- def update (self , input_pos , k_val , v_val ):
177
+ def update (self , input_pos , k_val , v_val , indices = None ):
165
178
"""
166
179
k_val, v_val: [B, H, S, D]
167
180
return: [B, H, S, D]
@@ -172,10 +185,12 @@ def update(self, input_pos, k_val, v_val):
172
185
v_val = v_val .transpose (1 , 2 )
173
186
174
187
if self .return_float_values :
175
- k_out , v_out = self ._update_and_return_float_values (input_pos , k_val , v_val )
188
+ k_out , v_out = self ._update_and_return_float_values (
189
+ input_pos , k_val , v_val , indices
190
+ )
176
191
else :
177
192
k_out , v_out = self ._update_and_return_quantized_values (
178
- input_pos , k_val , v_val
193
+ input_pos , k_val , v_val , indices
179
194
)
180
195
return k_out .transpose (1 , 2 ), v_out .transpose (1 , 2 )
181
196
@@ -277,14 +292,20 @@ def __init__(
277
292
)
278
293
279
294
def update (
280
- self , input_pos : torch .Tensor , k_val : torch .Tensor , v_val : torch .Tensor
295
+ self ,
296
+ input_pos : torch .Tensor ,
297
+ k_val : torch .Tensor ,
298
+ v_val : torch .Tensor ,
299
+ indices : Optional [torch .Tensor ] = None ,
281
300
) -> Tuple [torch .Tensor , torch .Tensor ]:
282
301
# input_pos: [S], k_val: [B, H, S, D]
283
302
k_val = k_val .transpose (1 , 2 )
284
303
v_val = v_val .transpose (1 , 2 )
285
304
start_pos = input_pos [0 ].item ()
286
- _ = torch .ops .llama .update_cache (k_val , self .k_cache , start_pos )
287
- _ = torch .ops .llama .update_cache (v_val , self .v_cache , start_pos )
305
+
306
+ _ = torch .ops .llama .update_cache (k_val , self .k_cache , start_pos , indices )
307
+ _ = torch .ops .llama .update_cache (v_val , self .v_cache , start_pos , indices )
308
+
288
309
return (
289
310
self .k_cache .transpose (1 , 2 ),
290
311
self .v_cache .transpose (1 , 2 ),
0 commit comments