test_qkv_proj_with_rope.py 11.3 KB
Newer Older
1
2
3
4
5
6
7
8
9
10
11
12
13
14
15
16
17
18
19
20
21
22
23
24
25
26
27
28
29
30
31
32
33
34
35
36
37
38
39
40
41
42
43
44
45
46
47
48
49
50
51
52
53
54
55
56
57
58
59
60
61
62
63
64
65
66
67
68
69
70
71
72
73
74
75
76
77
78
79
80
81
82
83
84
85
86
87
88
89
90
91
92
93
94
95
96
97
98
99
100
101
102
103
104
105
106
107
108
109
110
111
112
113
114
115
116
117
118
119
120
121
122
123
124
125
126
127
128
129
130
131
132
133
134
135
136
137
138
139
140
141
142
143
144
145
146
147
148
149
150
151
152
153
154
155
156
157
158
159
160
161
162
163
164
165
166
167
168
169
170
171
172
173
174
175
176
177
178
179
180
181
182
183
184
185
186
187
188
189
190
191
192
193
194
195
196
197
198
199
200
201
202
203
204
205
206
207
208
209
210
211
212
213
214
215
216
217
218
219
220
221
222
223
224
225
226
227
228
229
230
231
232
233
234
235
236
237
238
239
240
241
242
243
244
245
246
247
248
249
250
251
252
253
254
255
256
257
258
259
260
261
262
263
264
265
266
267
268
269
270
271
272
273
274
275
276
277
278
279
280
281
282
283
284
285
286
287
288
289
290
291
292
293
294
295
296
297
298
299
300
301
302
303
304
305
306
307
308
309
310
311
312
313
314
315
316
317
318
319
320
321
322
323
324
325
326
327
328
329
330
331
332
333
334
335
336
337
338
339
340
341
342
343
344
345
346
import unittest

import sgl_kernel
import torch
from utils import (
    convert_weight,
    native_w8a8_per_token_matmul,
    per_token_quant_int8,
    precision,
)

from sglang.srt.layers.rotary_embedding import _apply_rotary_emb
from sglang.test.test_utils import CustomTestCase

convert_weight_packed = torch.ops.sgl_kernel.convert_weight_packed
qkv_proj_with_rope = torch.ops.sgl_kernel.qkv_proj_with_rope
torch.manual_seed(0)
# constants
kv_lora_rank = 512
qk_head_dim = 192
qk_nope_head_dim = 128
qk_rope_head_dim = 64
rotary_dim = qk_rope_head_dim
num_heads = 22
q_lora_rank = 1536
hidden_size = 7168
B = 1
eps = 1e-6


def layernorm(x, weight, variance_epsilon=1e-6, residual=None):
    orig_dtype = x.dtype
    x = x.to(torch.float32)
    variance = x.pow(2).mean(dim=-1, keepdim=True)
    x = x * torch.rsqrt(variance + variance_epsilon)
    return (x * weight).to(orig_dtype)


def rotary_emb(q_pe, k_pe, pos, cos_sin_cache):
    orig_dtype = q_pe.dtype
    q_pe = q_pe.float()
    k_pe = k_pe.float()
    cos_sin_cache = cos_sin_cache.float()

    query_rot = q_pe[..., :rotary_dim]
    key_rot = k_pe[..., :rotary_dim]
    cos_sin = cos_sin_cache[pos]
    cos, sin = cos_sin.chunk(2, dim=-1)
    query_rot = _apply_rotary_emb(query_rot, cos, sin, False)
    key_rot = _apply_rotary_emb(key_rot, cos, sin, False)
    return query_rot.to(orig_dtype), key_rot.to(orig_dtype)


def native_torch(
    q_input,
    hidden_states,
    q_a_proj_weight,
    norm_weight1,
    q_b_proj_weight,
    w_kc,
    kv_a_proj_weight,
    norm_weight2,
    pos,
    cos_sin_cache,
):

    q = torch.matmul(hidden_states, q_a_proj_weight.t())
    q = layernorm(q, norm_weight1)
    q = torch.matmul(q, q_b_proj_weight.t()).view(-1, num_heads, qk_head_dim)

    q_nope, q_pe = q.split([qk_nope_head_dim, qk_rope_head_dim], dim=-1)
    q_nope_out = torch.bmm(q_nope.transpose(0, 1), w_kc)

    q_input[..., :kv_lora_rank] = q_nope_out.transpose(0, 1)
    latent_cache = torch.matmul(hidden_states, kv_a_proj_weight.t())
    v_input = latent_cache[..., :kv_lora_rank]
    v_input = layernorm(v_input.contiguous(), norm_weight2).unsqueeze(1)
    k_input = latent_cache.unsqueeze(1)
    k_input[..., :kv_lora_rank] = v_input
    k_pe = k_input[..., kv_lora_rank:]

    q_pe, k_pe = rotary_emb(q_pe, k_pe, pos, cos_sin_cache)
    q_input[..., kv_lora_rank:] = q_pe
    k_input[..., kv_lora_rank:] = k_pe

    return q_input, k_input, v_input


def native_torch_int8(
    q_input,
    hidden_states,
    w1_q,
    w1_s,
    norm_weight1,
    w2_q,
    w2_s,
    w_kc,
    w3_q,
    w3_s,
    norm_weight2,
    pos,
    cos_sin_cache,
):

    a_q, a_s = per_token_quant_int8(hidden_states)
    q = native_w8a8_per_token_matmul(a_q, w1_q, a_s, w1_s, None, torch.bfloat16)
    q = layernorm(q, norm_weight1)

    a_q, a_s = per_token_quant_int8(q)
    q = native_w8a8_per_token_matmul(a_q, w2_q, a_s, w2_s, None, torch.bfloat16).view(
        -1, num_heads, qk_head_dim
    )

    q_nope, q_pe = q.split([qk_nope_head_dim, qk_rope_head_dim], dim=-1)
    q_nope_out = torch.bmm(q_nope.transpose(0, 1), w_kc)

    q_input[..., :kv_lora_rank] = q_nope_out.transpose(0, 1)
    a_q, a_s = per_token_quant_int8(hidden_states)
    latent_cache = native_w8a8_per_token_matmul(
        a_q, w3_q, a_s, w3_s, None, torch.bfloat16
    )
    v_input = latent_cache[..., :kv_lora_rank]
    v_input = layernorm(v_input.contiguous(), norm_weight2).unsqueeze(1)
    k_input = latent_cache.unsqueeze(1)
    k_input[..., :kv_lora_rank] = v_input
    k_pe = k_input[..., kv_lora_rank:]

    q_pe, k_pe = rotary_emb(q_pe, k_pe, pos, cos_sin_cache)
    q_input[..., kv_lora_rank:] = q_pe
    k_input[..., kv_lora_rank:] = k_pe

    return q_input, k_input, v_input


class TestQKVProjWithROPE(CustomTestCase):
    def test_bf16_qkv_proj_with_rope(self):
        dtype = torch.bfloat16
        hidden_states = torch.randn(B, hidden_size, dtype=dtype) / hidden_size
        q_input = torch.empty(
            B, num_heads, kv_lora_rank + qk_rope_head_dim, dtype=dtype
        )
        q_a_proj_weight = torch.randn(q_lora_rank, hidden_size, dtype=dtype) * 0.1
        norm_weight1 = torch.randn(q_lora_rank, dtype=dtype)
        q_b_proj_weight = (
            torch.randn(num_heads * qk_head_dim, q_lora_rank, dtype=dtype) * 0.1
        )
        w_kc = torch.randn(num_heads, kv_lora_rank, qk_nope_head_dim, dtype=dtype) * 0.1
        kv_a_proj_weight = (
            torch.randn(kv_lora_rank + qk_rope_head_dim, hidden_size, dtype=dtype) * 0.1
        )
        norm_weight2 = torch.randn(kv_lora_rank, dtype=dtype)
        pos = torch.randint(10, 100, (B,))
        cos_sin_cache = torch.randn(100, rotary_dim, dtype=dtype)
        q_ref, k_ref, v_ref = native_torch(
            q_input,
            hidden_states,
            q_a_proj_weight,
            norm_weight1,
            q_b_proj_weight,
            w_kc.transpose(1, 2),
            kv_a_proj_weight,
            norm_weight2,
            pos,
            cos_sin_cache,
        )
        qa_packed = convert_weight_packed(q_a_proj_weight)
        qb_packed = convert_weight_packed(q_b_proj_weight)
        kva_packed = convert_weight_packed(kv_a_proj_weight)
        wkc_packed = convert_weight_packed(w_kc)

        q_out, k_out, v_out = qkv_proj_with_rope(
            hidden_states,
            qa_packed,
            qb_packed,
            kva_packed,
            wkc_packed,
            norm_weight1,
            norm_weight2,
            pos,
            cos_sin_cache,
            eps,
            False,
            False,
            None,
            None,
            None,
            True,
            None,
        )
        atol = rtol = precision[q_ref.dtype]
        self.assertTrue(torch.allclose(q_ref, q_out, atol=atol, rtol=rtol))
        self.assertTrue(torch.allclose(k_ref, k_out, atol=atol, rtol=rtol))
        self.assertTrue(torch.allclose(v_ref, v_out, atol=atol, rtol=rtol))

    def test_int8_qkv_proj_with_rope(self):
        dtype = torch.bfloat16
        hidden_states = torch.randn(B, hidden_size, dtype=dtype) / hidden_size
        q_input = torch.empty(
            B, num_heads, kv_lora_rank + qk_rope_head_dim, dtype=dtype
        )
        q_a_proj_weight = torch.randn(q_lora_rank, hidden_size, dtype=dtype) * 0.1
        norm_weight1 = torch.randn(q_lora_rank, dtype=dtype)
        q_b_proj_weight = (
            torch.randn(num_heads * qk_head_dim, q_lora_rank, dtype=dtype) * 0.1
        )
        w_kc = torch.randn(num_heads, kv_lora_rank, qk_nope_head_dim, dtype=dtype) * 0.1
        kv_a_proj_weight = (
            torch.randn(kv_lora_rank + qk_rope_head_dim, hidden_size, dtype=dtype) * 0.1
        )
        norm_weight2 = torch.randn(kv_lora_rank, dtype=dtype)
        pos = torch.randint(10, 100, (B,))
        cos_sin_cache = torch.randn(100, rotary_dim, dtype=dtype)

        w1_q, w1_s = per_token_quant_int8(q_a_proj_weight)
        w2_q, w2_s = per_token_quant_int8(q_b_proj_weight)
        w3_q, w3_s = per_token_quant_int8(kv_a_proj_weight)
        q_ref, k_ref, v_ref = native_torch_int8(
            q_input,
            hidden_states,
            w1_q,
            w1_s,
            norm_weight1,
            w2_q,
            w2_s,
            w_kc.transpose(1, 2),
            w3_q,
            w3_s,
            norm_weight2,
            pos,
            cos_sin_cache,
        )
        w1_q_packed = convert_weight_packed(w1_q)
        w2_q_packed = convert_weight_packed(w2_q)
        w3_q_packed = convert_weight_packed(w3_q)
        wkc_packed = convert_weight_packed(w_kc)
        q_out, k_out, v_out = qkv_proj_with_rope(
            hidden_states,
            w1_q_packed,
            w2_q_packed,
            w3_q_packed,
            wkc_packed,
            norm_weight1,
            norm_weight2,
            pos,
            cos_sin_cache,
            eps,
            True,
            False,
            w1_s,
            w2_s,
            w3_s,
            True,
            None,
        )
        atol = rtol = precision[q_ref.dtype]
        self.assertTrue(torch.allclose(q_ref, q_out, atol=atol, rtol=rtol))
        self.assertTrue(torch.allclose(k_ref, k_out, atol=atol, rtol=rtol))
        self.assertTrue(torch.allclose(v_ref, v_out, atol=atol, rtol=rtol))

    def test_fp8_qkv_proj_with_rope(self):
        dtype = torch.bfloat16
        hidden_states = torch.randn(B, hidden_size, dtype=dtype) / hidden_size
        q_input = torch.empty(
            B, num_heads, kv_lora_rank + qk_rope_head_dim, dtype=dtype
        )
        q_a_proj_weight = torch.randn(q_lora_rank, hidden_size, dtype=dtype) * 0.1
        norm_weight1 = torch.randn(q_lora_rank, dtype=dtype)
        q_b_proj_weight = (
            torch.randn(num_heads * qk_head_dim, q_lora_rank, dtype=dtype) * 0.1
        )
        w_kc = torch.randn(num_heads, kv_lora_rank, qk_nope_head_dim, dtype=dtype) * 0.1
        kv_a_proj_weight = (
            torch.randn(kv_lora_rank + qk_rope_head_dim, hidden_size, dtype=dtype) * 0.1
        )
        norm_weight2 = torch.randn(kv_lora_rank, dtype=dtype)
        pos = torch.randint(10, 100, (B,))
        cos_sin_cache = torch.randn(100, rotary_dim, dtype=dtype)

        scale_block_size_N = 128
        scale_block_size_K = 128
        fp8_q_a_proj_weight, q_a_proj_weight_scale_inv, q_a_proj_weight_dq = (
            convert_weight(
                q_a_proj_weight,
                [scale_block_size_N, scale_block_size_K],
                torch.bfloat16,
            )
        )
        fp8_q_b_proj_weight, q_b_proj_weight_scale_inv, q_b_proj_weight_dq = (
            convert_weight(
                q_b_proj_weight,
                [scale_block_size_N, scale_block_size_K],
                torch.bfloat16,
            )
        )
        (
            fp8_kv_a_proj_with_mqa_weight,
            kv_a_proj_with_mqa_weight_scale_inv,
            kv_a_proj_with_mqa_weight_dq,
        ) = convert_weight(
            kv_a_proj_weight, [scale_block_size_N, scale_block_size_K], torch.bfloat16
        )
        q_ref, k_ref, v_ref = native_torch(
            q_input,
            hidden_states,
            q_a_proj_weight_dq,
            norm_weight1,
            q_b_proj_weight_dq,
            w_kc.transpose(1, 2),
            kv_a_proj_with_mqa_weight_dq,
            norm_weight2,
            pos,
            cos_sin_cache,
        )
        fp8_q_a_proj_weight = convert_weight_packed(fp8_q_a_proj_weight)
        fp8_q_b_proj_weight = convert_weight_packed(fp8_q_b_proj_weight)
        fp8_kv_a_proj_with_mqa_weight = convert_weight_packed(
            fp8_kv_a_proj_with_mqa_weight
        )
        w_kc = convert_weight_packed(w_kc)
        q_out, k_out, v_out = qkv_proj_with_rope(
            hidden_states,
            fp8_q_a_proj_weight,
            fp8_q_b_proj_weight,
            fp8_kv_a_proj_with_mqa_weight,
            w_kc,
            norm_weight1,
            norm_weight2,
            pos,
            cos_sin_cache,
            eps,
            False,
            True,
            q_a_proj_weight_scale_inv.float(),
            q_b_proj_weight_scale_inv.float(),
            kv_a_proj_with_mqa_weight_scale_inv.float(),
            True,
            [scale_block_size_N, scale_block_size_K],
        )
        atol = rtol = precision[q_ref.dtype]
        self.assertTrue(torch.allclose(q_ref, q_out, atol=atol, rtol=rtol))
        self.assertTrue(torch.allclose(k_ref, k_out, atol=atol, rtol=rtol))
        self.assertTrue(torch.allclose(v_ref, v_out, atol=atol, rtol=rtol))


if __name__ == "__main__":
    unittest.main()