test_triton_flash_attention.py 18.8 KB
Newer Older
1
# SPDX-License-Identifier: Apache-2.0
2
# SPDX-FileCopyrightText: Copyright contributors to the vLLM project
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
347
348
349
350
351
352
353
354
355
356
357
358
359
360
361
362
363
364
365
366
367
368
369
370
371
372
373
374
375
376
377
378
379
380
381
382
383
384
385
386
387
388
389
390
391
392
393
394
395
396
397
398
399
400
401
402
403
404
405
406
407
408
409
410
411
412
413
414
415
416
417
418
419
420
421
422
423
424
425
426
427
428
429
430
431
432
433
434
435
436
437
438
439
440
441
442
443
444
445
446
447
448
449
450
451
452
453
454
455
456
457
458
459
460
461
462
463
464
465
466
467
468
469
470
471
472
473
474
475
476
477
478
479
480
481
482
483
484
485
486
487
488
489
490
491
492
493
494
495
496
497
498
499
500
"""Tests for the triton_flash_attention kernel

Run `pytest tests/kernels/test_triton_flash_attention.py`.
"""
import pytest
import torch

from vllm.attention.ops.triton_flash_attention import (SUPPORTED_LAYOUTS,
                                                       MetaData,
                                                       compute_alibi_tensor,
                                                       scale_fp8,
                                                       triton_attention_rocm)
from vllm.platforms import current_platform


class ReferenceAttention:

    def __init__(self, Z, HQ, HK, N_CTX_Q, N_CTX_K, D_HEAD, use_alibi, dtype,
                 input_metadata):
        self.Z = Z
        self.HQ = HQ
        self.HK = HK
        self.N_CTX_Q = N_CTX_Q
        self.N_CTX_K = N_CTX_K
        self.D_HEAD = D_HEAD
        self.use_alibi = use_alibi
        self.dtype = dtype
        self.input_metadata = input_metadata

    def fwd(self, q, k, v):
        scores = torch.einsum('bhqd,bhkd->bhqk', q,
                              k).float() * self.input_metadata.sm_scale
        if self.input_metadata.causal:
            mask = torch.tril(torch.ones(self.N_CTX_Q,
                                         self.N_CTX_K,
                                         device="cuda"),
                              diagonal=self.N_CTX_K - self.N_CTX_Q)
            scores[:, :, mask == 0] = float("-inf")

        if self.input_metadata.bias is not None:
            scores += self.input_metadata.bias

        if self.use_alibi:
            scores += compute_alibi_tensor(self.input_metadata.alibi_slopes,
                                           self.N_CTX_Q, self.N_CTX_K)

        p = torch.softmax(scores, dim=-1)
        if self.input_metadata.causal:
            # If N_CTX_Q > N_CTX_K, there's at least one row of all -infs going
            # into softmax. This creates a row of NaNs as -inf - -inf == NaN.
            # So we fix this by converting the NaNs to 0s, which is what they
            # should be out of the softmax.
            nan_mask = torch.isnan(p)
            p[nan_mask == 1] = 0
        ref_out = torch.einsum('bhqk,bhkd->bhqd', p.to(self.dtype), v)
        # compare
        if self.input_metadata.layout == 'bshd':
            ref_out = ref_out.transpose(1, 2).clone()
        return ref_out

    def fwd_fp8(self, q_quantized, k_quantized, v_quantized):
        q = (q_quantized.to(torch.float16) * self.input_metadata.q_descale).to(
            self.dtype)
        k = (k_quantized.to(torch.float16) * self.input_metadata.k_descale).to(
            self.dtype)
        v = (v_quantized.to(torch.float16) * self.input_metadata.v_descale).to(
            self.dtype)
        result = self.fwd(q, k, v)
        if self.input_metadata.o_scale is not None:
            result, _ = scale_fp8(result, self.input_metadata.o_scale)
        return result

    def fwd_fp8_kv(self, q, k_quantized, v_quantized):
        k_descale, v_descale = (self.input_metadata.k_descale,
                                self.input_metadata.v_descale)
        k_dequantized = (k_quantized.to(torch.float32) *
                         k_descale.to(torch.float32)).to(self.dtype)
        v_dequantized = (v_quantized.to(torch.float32) *
                         v_descale.to(torch.float32)).to(self.dtype)
        return self.fwd(q, k_dequantized, v_dequantized)

    def varlen_fwd(self, q, k, v, is_mqa=False):
        ref_out = torch.empty_like(q)
        if is_mqa:
            # Make KV look like HQ/HK "groups" of HK. Later, we will reshape so
            # the size aligns with Q.
            k_ref = k.view(k.shape[0], k.shape[1], 1,
                           k.shape[2]).expand(-1, -1, self.HQ // self.HK, -1)
            v_ref = v.view(v.shape[0], v.shape[1], 1,
                           v.shape[2]).expand(-1, -1, self.HQ // self.HK, -1)
        else:
            k_ref = k
            v_ref = v

        for i in range(0, self.input_metadata.num_contexts):
            start_q, start_k = self.input_metadata.cu_seqlens_q[
                i], self.input_metadata.cu_seqlens_k[i]
            end_q, end_k = self.input_metadata.cu_seqlens_q[
                i + 1], self.input_metadata.cu_seqlens_k[i + 1]
            k_curr = k_ref[start_k:end_k]
            v_curr = v_ref[start_k:end_k]
            if is_mqa:
                k_curr = k_curr.reshape(k_curr.shape[0], -1, k_curr.shape[3])
                v_curr = v_curr.reshape(v_curr.shape[0], -1, v_curr.shape[3])
            scores = torch.einsum('qhd,khd->qhk', q[start_q:end_q],
                                  k_curr).float()
            p = torch.softmax(scores * self.input_metadata.sm_scale,
                              dim=-1).half()
            ref_out[start_q:end_q] = torch.einsum('qhk,khd->qhd', p, v_curr)
        return ref_out


def quantize_input(q, k, v, fp8_kv=False, use_o_scale=False):
    q_descale = None
    if not fp8_kv:
        q, q_descale = scale_fp8(q)
    k, k_descale = scale_fp8(k)
    v, v_descale = scale_fp8(v)

    # In real world use case, the p scale would be a parameter trained by the
    # model.
    p_scale = None

    o_scale = torch.rand(1, device="cuda",
                         requires_grad=False) if use_o_scale else None

    return q, k, v, q_descale, k_descale, v_descale, p_scale, o_scale


def input_helper(
    Z,
    HQ,
    HK,
    N_CTX_Q,
    N_CTX_K,
    D_HEAD,
    dtype,
    layout=None,
    use_alibi=None,
    causal=None,
    is_fp8=False,
    fp8_kv=False,
    use_o_scale=False,
    use_bias=False,
):
    assert layout in SUPPORTED_LAYOUTS, "Got unsupported layout."

    current_platform.seed_everything(0)

    # Initialize q, k, v
    if layout == 'bhsd':
        q_tensor_shape = (Z, HQ, N_CTX_Q, D_HEAD)
        k_tensor_shape = (Z, HK, N_CTX_K, D_HEAD)
    elif layout == 'bshd':
        q_tensor_shape = (Z, N_CTX_Q, HQ, D_HEAD)
        k_tensor_shape = (Z, N_CTX_K, HK, D_HEAD)

    if use_alibi:
        # for n heads the set of slopes is the geometric sequence that starts
        # 2^(-8/n)
        alibi_slopes = torch.tensor(
            [2**(-8 / HQ * i) for i in range(1, HQ + 1)],
            dtype=torch.float32,
            device="cuda").repeat(Z, 1)
    else:
        alibi_slopes = None

    if use_bias:
        bias = torch.randn((1, HQ, N_CTX_Q, N_CTX_K),
                           dtype=dtype,
                           device="cuda",
                           requires_grad=False)
    else:
        bias = None

    q = torch.randn(q_tensor_shape,
                    dtype=dtype,
                    device="cuda",
                    requires_grad=False)
    k = torch.randn(k_tensor_shape,
                    dtype=dtype,
                    device="cuda",
                    requires_grad=False)
    v = torch.randn(k_tensor_shape,
                    dtype=dtype,
                    device="cuda",
                    requires_grad=False)

    if is_fp8:
        (q, k, v, q_descale, k_descale, v_descale, p_scale,
         o_scale) = quantize_input(q,
                                   k,
                                   v,
                                   use_o_scale=use_o_scale,
                                   fp8_kv=fp8_kv)
    else:
        q_descale = k_descale = v_descale = p_scale = o_scale = None

    input_metadata = MetaData(sm_scale=D_HEAD**-0.5,
                              max_seqlens_q=N_CTX_Q,
                              max_seqlens_k=N_CTX_K,
                              layout=layout,
                              alibi_slopes=alibi_slopes,
                              alibi_batch=Z,
                              alibi_nheads=HQ,
                              q_descale=q_descale,
                              k_descale=k_descale,
                              v_descale=v_descale,
                              p_scale=p_scale,
                              o_scale=o_scale,
                              bias=bias,
                              seqlen_q=N_CTX_Q,
                              seqlen_k=N_CTX_K)
    return q, k, v, input_metadata


def varlen_input_helper(Z,
                        HQ,
                        HK,
                        N_CTX_Q,
                        N_CTX_K,
                        D_HEAD,
                        dtype,
                        equal_seqlens=False):
    current_platform.seed_everything(0)

    # Random sequence lengths. Using N_CTX as kind of max of sum of individual
    # seqs
    if not equal_seqlens:
        max_seqlens_q = N_CTX_Q // Z
        max_seqlens_k = N_CTX_K // Z
        seqlens_q = torch.randint(1,
                                  max_seqlens_q + 1, (Z, ),
                                  dtype=torch.int32)
        seqlens_k = torch.randint(1,
                                  max_seqlens_k + 1, (Z, ),
                                  dtype=torch.int32)
    else:
        seqlens_q = torch.full((Z, ), N_CTX_Q // Z)
        seqlens_k = torch.full((Z, ), N_CTX_K // Z)

    # Calculate cumulative sequence lengths
    cu_seqlens_q = torch.cat([
        torch.tensor([0], dtype=torch.int32),
        seqlens_q.cumsum(dim=0, dtype=torch.int32)
    ])
    cu_seqlens_k = torch.cat([
        torch.tensor([0], dtype=torch.int32),
        seqlens_k.cumsum(dim=0, dtype=torch.int32)
    ])
    cu_seqlens_q = cu_seqlens_q.to(device="cuda")
    cu_seqlens_k = cu_seqlens_k.to(device="cuda")

    # Initialize q, k, v with variable lengths
    total_q = cu_seqlens_q[-1].item()
    total_k = cu_seqlens_k[-1].item()
    q = torch.randn((total_q, HQ, D_HEAD), dtype=dtype,
                    device="cuda").normal_(mean=0., std=0.5).requires_grad_()
    k = torch.randn((total_k, HK, D_HEAD), dtype=dtype,
                    device="cuda").normal_(mean=0., std=0.5).requires_grad_()
    v = torch.randn((total_k, HK, D_HEAD), dtype=dtype,
                    device="cuda").normal_(mean=0., std=0.5).requires_grad_()
    sm_scale = D_HEAD**-0.5
    input_metadata = MetaData(sm_scale=sm_scale)
    input_metadata.set_varlen_params(cu_seqlens_q, cu_seqlens_k)
    return q, k, v, input_metadata


@pytest.mark.parametrize('Z, HQ, HK, N_CTX_Q, N_CTX_K, D_HEAD', [
    (1, 48, 12, 1, 1, 64),
    (4, 4, 4, 128, 128, 65),
    (16, 48, 48, 1, 1, 128),
    (64, 48, 24, 3, 3, 128),
    (4, 4, 4, 113, 123, 1),
])
@pytest.mark.parametrize('causal', [True, False])
@pytest.mark.parametrize('use_alibi', [True, False])
@pytest.mark.parametrize('layout', ['bshd'])
def test_op_fwd(Z,
                HQ,
                HK,
                N_CTX_Q,
                N_CTX_K,
                D_HEAD,
                causal,
                use_alibi,
                layout,
                dtype=torch.float16):
    current_platform.seed_everything(0)
    q, k, v, input_metadata = input_helper(Z, HQ, HK, N_CTX_Q, N_CTX_K, D_HEAD,
                                           dtype, layout, use_alibi, causal)

    o = torch.empty_like(q)

    # triton implementation
    tri_out, _ = triton_attention_rocm(q, k, v, o, input_metadata)

    # Transpose here if layout is bshd so we have same reference code for all
    # layouts
    if layout == 'bshd':
        q = q.transpose(1, 2).clone()
        k = k.transpose(1, 2).clone()
        v = v.transpose(1, 2).clone()
    # Replicate K and V if using MQA/GQA
    if HQ != HK:
        k = k.view(k.shape[0], k.shape[1], -1, k.shape[2],
                   k.shape[3]).expand(-1, -1, HQ // HK, -1,
                                      -1).reshape(k.shape[0], -1, k.shape[2],
                                                  k.shape[3])
        v = v.view(v.shape[0], v.shape[1], -1, v.shape[2],
                   v.shape[3]).expand(-1, -1, HQ // HK, -1,
                                      -1).reshape(v.shape[0], -1, v.shape[2],
                                                  v.shape[3])

    ref_impl = ReferenceAttention(Z, HQ, HK, N_CTX_Q, N_CTX_K, D_HEAD,
                                  use_alibi, dtype, input_metadata)
    ref_out = ref_impl.fwd(q, k, v)

    torch.testing.assert_close(ref_out, tri_out, atol=2e-2, rtol=2e-2)


@pytest.mark.parametrize('Z, H, N_CTX_Q, N_CTX_K, D_HEAD', [
    (4, 48, 1, 1, 64),
    (4, 48, 1, 1, 128),
    (4, 48, 3, 3, 128),
    (4, 4, 128, 128, 65),
])
@pytest.mark.parametrize('causal', [True, False])
@pytest.mark.parametrize('layout', ['bhsd'])
@pytest.mark.parametrize('use_o_scale', [True, False])
@pytest.mark.skipif(torch.cuda.get_device_capability() < (9, 0),
                    reason="Triton FP8 requires CUDA 9.0 or higher")
def test_op_fwd_fp8(Z,
                    H,
                    N_CTX_Q,
                    N_CTX_K,
                    D_HEAD,
                    causal,
                    layout,
                    use_o_scale,
                    dtype=torch.float32):
    current_platform.seed_everything(0)

    # Disable grad to save memory it won't run into OOM on CI machine.
    # q, k, v, input_metadata = input_helper(Z, H, H, N_CTX_Q, N_CTX_K, D_HEAD,
    # dtype, layout)

    q_quantized, k_quantized, v_quantized, input_metadata = input_helper(
        Z,
        H,
        H,
        N_CTX_Q,
        N_CTX_K,
        D_HEAD,
        dtype,
        causal=causal,
        layout=layout,
        is_fp8=True,
        use_o_scale=use_o_scale)

    o = torch.empty_like(q_quantized) if use_o_scale else None

    tri_out, _ = triton_attention_rocm(q_quantized, k_quantized, v_quantized,
                                       o, input_metadata)

    ref_impl = ReferenceAttention(Z, H, H, N_CTX_Q, N_CTX_K, D_HEAD, False,
                                  dtype, input_metadata)
    ref_out = ref_impl.fwd_fp8(q_quantized, k_quantized, v_quantized)

    # compare
    torch.testing.assert_close(ref_out.to(torch.float32),
                               tri_out.to(torch.float32),
                               atol=7e-2,
                               rtol=2e-1)


@pytest.mark.parametrize('Z, H, N_CTX_Q, N_CTX_K, D_HEAD', [
    (4, 48, 1, 1, 64),
    (4, 48, 1, 1, 128),
    (4, 48, 3, 3, 128),
    (4, 4, 128, 128, 65),
    (4, 4, 113, 123, 1),
])
@pytest.mark.parametrize('causal', [True, False])
@pytest.mark.parametrize('layout', ['bhsd'])
def test_op_fwd_fp8_kv(Z,
                       H,
                       N_CTX_Q,
                       N_CTX_K,
                       D_HEAD,
                       causal,
                       layout,
                       dtype=torch.float32):
    current_platform.seed_everything(0)

    q, k_quantized, v_quantized, input_metadata = input_helper(Z,
                                                               H,
                                                               H,
                                                               N_CTX_Q,
                                                               N_CTX_K,
                                                               D_HEAD,
                                                               dtype,
                                                               causal=causal,
                                                               layout=layout,
                                                               is_fp8=True,
                                                               fp8_kv=True)

    o = torch.empty_like(q)

    tri_out, _ = triton_attention_rocm(q, k_quantized, v_quantized, o,
                                       input_metadata)

    ref_impl = ReferenceAttention(Z, H, H, N_CTX_Q, N_CTX_K, D_HEAD, False,
                                  dtype, input_metadata)
    ref_out = ref_impl.fwd_fp8_kv(q, k_quantized, v_quantized)

    torch.testing.assert_close(ref_out, tri_out, atol=3e-2, rtol=8e-1)


@pytest.mark.parametrize('Z, H, N_CTX_Q, N_CTX_K, D_HEAD', [
    (4, 48, 1, 1, 64),
    (4, 48, 1, 1, 128),
    (4, 48, 3, 3, 128),
    (4, 4, 128, 128, 65),
])
@pytest.mark.parametrize('causal', [True, False])
@pytest.mark.parametrize('use_bias', [True])
@pytest.mark.parametrize('dtype', [torch.bfloat16])
def test_op_fwd_bias(Z, H, N_CTX_Q, N_CTX_K, D_HEAD, causal, use_bias, dtype):
    current_platform.seed_everything(0)
    q, k, v, input_metadata = input_helper(Z,
                                           H,
                                           H,
                                           N_CTX_Q,
                                           N_CTX_K,
                                           D_HEAD,
                                           dtype,
                                           layout='bhsd',
                                           causal=causal,
                                           use_bias=use_bias)
    o = torch.empty_like(q)

    # triton implementation
    tri_out, _ = triton_attention_rocm(q, k, v, o, input_metadata)

    ref_impl = ReferenceAttention(Z, H, H, N_CTX_Q, N_CTX_K, D_HEAD, False,
                                  dtype, input_metadata)
    ref_out = ref_impl.fwd(q, k, v)

    # compare
    torch.testing.assert_close(ref_out, tri_out, atol=2e-2, rtol=2e-2)


# NOTE: Uses thd layout, so also tests thd.
@pytest.mark.parametrize('Z, H, N_CTX, D_HEAD', [(1, 48, 256, 64),
                                                 (4, 48, 512, 64),
                                                 (16, 48, 512, 64),
                                                 (64, 48, 128, 128)])
@pytest.mark.parametrize('causal', [True, False])
def test_op_varlen_fwd(Z, H, N_CTX, D_HEAD, causal, dtype=torch.float16):

    q, k, v, input_metadata = varlen_input_helper(Z, H, H, N_CTX, N_CTX,
                                                  D_HEAD, dtype)

    tri_out = torch.empty_like(q)
    triton_attention_rocm(q, k, v, tri_out, input_metadata)

    ref_impl = ReferenceAttention(Z, H, H, N_CTX, N_CTX, D_HEAD, False, dtype,
                                  input_metadata)
    ref_out = ref_impl.varlen_fwd(q, k, v, is_mqa=False)

    torch.testing.assert_close(ref_out, tri_out, atol=2e-2, rtol=2e-2)


# NOTE: Uses thd layout, so also tests thd.
@pytest.mark.parametrize('Z, HQ, HK, N_CTX, D_HEAD', [(2, 48, 24, 128, 64),
                                                      (4, 48, 12, 256, 64),
                                                      (4, 48, 4, 512, 64),
                                                      (4, 64, 16, 128, 128)])
@pytest.mark.parametrize('causal', [False])
def test_op_varlen_mqa_fwd(Z,
                           HQ,
                           HK,
                           N_CTX,
                           D_HEAD,
                           causal,
                           dtype=torch.float16):
    q, k, v, input_metadata = varlen_input_helper(Z, HQ, HK, N_CTX, N_CTX,
                                                  D_HEAD, dtype)

    tri_out = torch.empty_like(q)
    triton_attention_rocm(q, k, v, tri_out, input_metadata)

    ref_impl = ReferenceAttention(Z, HQ, HK, N_CTX, N_CTX, D_HEAD, False,
                                  dtype, input_metadata)
    ref_out = ref_impl.varlen_fwd(q, k, v, is_mqa=True)

    torch.testing.assert_close(ref_out, tri_out, atol=2e-2, rtol=2e-2)