test_attention.py 13.6 KB
Newer Older
1
import random
2
from typing import List, Optional, Tuple
3

4
import pytest
5
import torch
6
7
from xformers import ops as xops
from xformers.ops.fmha.attn_bias import BlockDiagonalCausalMask
8

9
from tests.kernels.utils import opcheck
10
from vllm import _custom_ops as ops
11
from vllm.utils import get_max_shared_memory_bytes, is_hip
12

13
14
from .allclose_default import get_default_atol, get_default_rtol

15
16
17
18
FLOAT32_BYTES = torch.finfo(torch.float).bits // 8
# This will change depending on the compute capability.
# - 512 as a buffer
MAX_SEQ_LEN = get_max_shared_memory_bytes() // FLOAT32_BYTES - 512
19
20
21
# There may not be enough gpu memory due to large NUM_BLOCKS.
# Reduce NUM_BLOCKS when it happens.
NUM_BLOCKS = 4321  # Arbitrary values for testing
22
PARTITION_SIZE = 512
23
24
25
# flshattF and tritonflashattF supported: {torch.float16, torch.bfloat16}
DTYPES = [torch.half, torch.bfloat16, torch.float
          ] if not is_hip() else [torch.half, torch.bfloat16]
26
NUM_GEN_SEQS = [7]  # Arbitrary values for testing
27
NUM_PREFILL_SEQS = [3]  # Arbitrary values for testing
28
NUM_HEADS = [(40, 40), (64, 8)]  # Arbitrary values for testing
29
30
31

# FlashAttention forward only supports head dimension at most 128
# https://github.com/ROCmSoftwarePlatform/flash-attention/blob/3d2b6f5d037782cc2c906909a46fb7e2e1b48b25/csrc/flash_attn_rocm/flash_api.cpp#L62
Joe's avatar
Joe committed
32
HEAD_SIZES = [64, 80, 96, 112, 120, 128, 192, 256
33
34
              ] if not is_hip() else [64, 80, 96, 112, 128]

35
BLOCK_SIZES = [16, 32]
36
USE_ALIBI = [False, True]
37
KV_CACHE_DTYPE = ["auto", "fp8"]
38
SEEDS = [0]
39
40
41
CUDA_DEVICES = [
    f"cuda:{i}" for i in range(1 if torch.cuda.device_count() == 1 else 2)
]
42

43
44
45
46
47
48
49
50

def ref_masked_attention(
    query: torch.Tensor,
    key: torch.Tensor,
    value: torch.Tensor,
    scale: float,
    attn_mask: Optional[torch.Tensor] = None,
) -> torch.Tensor:
51
    attn_weights = scale * torch.einsum("qhd,khd->hqk", query, key).float()
52
    if attn_mask is not None:
53
54
55
        attn_weights = attn_weights + attn_mask.float()
    attn_weights = torch.softmax(attn_weights, dim=-1).to(value.dtype)
    out = torch.einsum("hqk,khd->qhd", attn_weights, value)
56
57
58
59
60
61
    return out


def ref_single_query_cached_kv_attention(
    output: torch.Tensor,
    query: torch.Tensor,
62
    num_queries_per_kv: int,
63
64
65
    key_cache: torch.Tensor,
    value_cache: torch.Tensor,
    block_tables: torch.Tensor,
66
    seq_lens: torch.Tensor,
67
68
    scale: float,
    alibi_slopes: Optional[torch.Tensor],
69
) -> None:
70
71
    num_query_heads = query.shape[1]
    num_kv_heads = value_cache.shape[1]
72
73
    head_size = value_cache.shape[2]
    block_size = value_cache.shape[3]
74
    num_seqs = query.shape[0]
75

76
77
    block_tables_lst = block_tables.cpu().tolist()
    seq_lens_lst = seq_lens.cpu().tolist()
78
    for i in range(num_seqs):
79
        q = query[i].unsqueeze(0)
80
81
        block_table = block_tables_lst[i]
        seq_len = int(seq_lens_lst[i])
82

83
84
        keys_lst: List[torch.Tensor] = []
        values_lst: List[torch.Tensor] = []
85
        for j in range(seq_len):
86
87
88
89
            block_number = int(block_table[j // block_size])
            block_offset = j % block_size

            k = key_cache[block_number, :, :, block_offset, :]
90
            k = k.reshape(num_kv_heads, head_size)
91
            keys_lst.append(k)
92
93

            v = value_cache[block_number, :, :, block_offset]
94
95
96
            values_lst.append(v)
        keys = torch.stack(keys_lst, dim=0)
        values = torch.stack(values_lst, dim=0)
97
98
99
100
101
102
103
104
        if num_queries_per_kv > 1:
            # Handle MQA and GQA
            keys = torch.repeat_interleave(keys, num_queries_per_kv, dim=1)
            values = torch.repeat_interleave(values, num_queries_per_kv, dim=1)

        alibi_bias = None
        if alibi_slopes is not None:
            # Create the ALiBi bias used in the paged attention kernel.
105
106
            position_ids = torch.arange(seq_len).int()
            alibi_bias = (position_ids - seq_len + 1).float()
107
108
109
110
111
            alibi_bias = alibi_slopes.view(-1, 1, 1) * alibi_bias.view(
                1, 1, -1)

        out = ref_masked_attention(q, keys, values, scale, alibi_bias)
        out = out.view(num_query_heads, head_size)
112
113
114
        output[i].copy_(out, non_blocking=True)


115
@pytest.mark.parametrize("version", ["v1", "v2"])
116
117
118
119
120
121
@pytest.mark.parametrize("num_seqs", NUM_GEN_SEQS)
@pytest.mark.parametrize("num_heads", NUM_HEADS)
@pytest.mark.parametrize("head_size", HEAD_SIZES)
@pytest.mark.parametrize("use_alibi", USE_ALIBI)
@pytest.mark.parametrize("block_size", BLOCK_SIZES)
@pytest.mark.parametrize("dtype", DTYPES)
122
@pytest.mark.parametrize("kv_cache_dtype", KV_CACHE_DTYPE)
123
@pytest.mark.parametrize("seed", SEEDS)
124
@pytest.mark.parametrize("device", CUDA_DEVICES)
125
def test_paged_attention(
126
    kv_cache_factory,
127
    version: str,
128
129
    num_seqs: int,
    num_heads: Tuple[int, int],
130
    head_size: int,
131
    use_alibi: bool,
132
133
    block_size: int,
    dtype: torch.dtype,
134
    kv_cache_dtype: str,
135
    seed: int,
136
    device: str,
137
) -> None:
Joe's avatar
Joe committed
138
139
    if kv_cache_dtype == "fp8" and head_size % 16:
        pytest.skip()
140
141
    random.seed(seed)
    torch.random.manual_seed(seed)
142
143
144
    if torch.cuda.is_available():
        torch.cuda.manual_seed(seed)
    torch.set_default_device(device)
145
146
    scale = float(1.0 / (head_size**0.5))
    num_query_heads, num_kv_heads = num_heads
147
    query = torch.empty(num_seqs, num_query_heads, head_size, dtype=dtype)
148
149
150
151
152
153
    query.uniform_(-scale, scale)

    assert num_query_heads % num_kv_heads == 0
    num_queries_per_kv = num_query_heads // num_kv_heads
    alibi_slopes = None
    if use_alibi:
154
        alibi_slopes = torch.randn(num_query_heads, dtype=torch.float)
155

156
157
158
159
    seq_lens = [random.randint(1, MAX_SEQ_LEN) for _ in range(num_seqs)]
    seq_lens[-1] = MAX_SEQ_LEN
    max_seq_len = max(seq_lens)
    seq_lens = torch.tensor(seq_lens, dtype=torch.int)
160

161
    # Create the block tables.
162
    max_num_blocks_per_seq = (max_seq_len + block_size - 1) // block_size
163
    block_tables_lst: List[List[int]] = []
164
    for _ in range(num_seqs):
165
        block_table = [
166
            random.randint(0, NUM_BLOCKS - 1)
167
168
            for _ in range(max_num_blocks_per_seq)
        ]
169
170
171
        block_tables_lst.append(block_table)

    block_tables = torch.tensor(block_tables_lst, dtype=torch.int)
172

173
174
    # Create the KV caches.
    key_caches, value_caches = kv_cache_factory(NUM_BLOCKS, block_size, 1,
175
176
                                                num_kv_heads, head_size,
                                                kv_cache_dtype, dtype, seed,
177
                                                device)
178
    key_cache, value_cache = key_caches[0], value_caches[0]
Tao Peng's avatar
Tao Peng committed
179

180
    # Using default kv_scale
181
    k_scale = v_scale = 1.0
182

183
184
    # Call the paged attention kernel.
    output = torch.empty_like(query)
185
    if version == "v1":
186
        ops.paged_attention_v1(
187
188
189
190
            output,
            query,
            key_cache,
            value_cache,
191
            num_kv_heads,
192
193
            scale,
            block_tables,
194
            seq_lens,
195
            block_size,
196
            max_seq_len,
197
            alibi_slopes,
198
            kv_cache_dtype,
199
200
            k_scale,
            v_scale,
201
        )
202
203
204
205
206
207
208

        opcheck(torch.ops._C.paged_attention_v1,
                (output, query, key_cache, value_cache, num_kv_heads, scale,
                 block_tables, seq_lens, block_size, max_seq_len, alibi_slopes,
                 kv_cache_dtype, k_scale, v_scale, 0, 0, 0, 64, 0),
                cond=(head_size == HEAD_SIZES[0]))

209
    elif version == "v2":
210
        num_partitions = ((max_seq_len + PARTITION_SIZE - 1) // PARTITION_SIZE)
211
212
213
214
215
216
217
218
219
220
221
        assert PARTITION_SIZE % block_size == 0
        num_seqs, num_heads, head_size = output.shape
        tmp_output = torch.empty(
            size=(num_seqs, num_heads, num_partitions, head_size),
            dtype=output.dtype,
        )
        exp_sums = torch.empty(
            size=(num_seqs, num_heads, num_partitions),
            dtype=torch.float32,
        )
        max_logits = torch.empty_like(exp_sums)
222
        ops.paged_attention_v2(
223
224
225
226
227
228
229
            output,
            exp_sums,
            max_logits,
            tmp_output,
            query,
            key_cache,
            value_cache,
230
            num_kv_heads,
231
232
            scale,
            block_tables,
233
            seq_lens,
234
            block_size,
235
            max_seq_len,
236
            alibi_slopes,
237
            kv_cache_dtype,
238
239
            k_scale,
            v_scale,
240
        )
241
242
243
244
245
246
247
248

        opcheck(torch.ops._C.paged_attention_v2,
                (output, exp_sums, max_logits, tmp_output, query, key_cache,
                 value_cache, num_kv_heads, scale, block_tables, seq_lens,
                 block_size, max_seq_len, alibi_slopes, kv_cache_dtype,
                 k_scale, v_scale, 0, 0, 0, 64, 0),
                cond=(head_size == HEAD_SIZES[0]))

249
    else:
250
        raise AssertionError(f"Unknown version: {version}")
251

252
    # Run the reference implementation.
253
    if kv_cache_dtype == "fp8":
254
255
256
257
258
259
        # Convert cache data back to dtype.
        x = 16 // torch.tensor([], dtype=dtype).element_size()
        key_cache_shape = (NUM_BLOCKS, num_kv_heads, head_size // x,
                           block_size, x)
        dequantized_key_cache = torch.empty(size=key_cache_shape,
                                            dtype=dtype,
260
                                            device=device)
261
        ops.convert_fp8(dequantized_key_cache, key_cache)
262
263
264
265
266
        key_cache = dequantized_key_cache

        value_cache_shape = value_cache.shape
        dequantized_value_cache = torch.empty(size=value_cache_shape,
                                              dtype=dtype,
267
                                              device=device)
268
        ops.convert_fp8(dequantized_value_cache, value_cache)
269
270
        value_cache = dequantized_value_cache

271
272
273
274
    ref_output = torch.empty_like(query)
    ref_single_query_cached_kv_attention(
        ref_output,
        query,
275
        num_queries_per_kv,
276
277
278
        key_cache,
        value_cache,
        block_tables,
279
        seq_lens,
280
281
        scale,
        alibi_slopes,
282
    )
283
284
285
286

    # NOTE(woosuk): Due to the kernel-level differences in the two
    # implementations, there is a small numerical difference in the two
    # outputs. Thus, we use a relaxed tolerance for the test.
287
288
289
    atol = get_default_atol(output) if is_hip() else 1e-3
    rtol = get_default_rtol(output) if is_hip() else 1e-5

290
291
    # NOTE(zhaoyang): FP8 KV Cache will introduce quantization error,
    # so we use a relaxed tolerance for the test.
292
293
    atol, rtol = 1e-3, 1e-5
    if kv_cache_dtype == "fp8":
294
        atol, rtol = 1e-2, 1e-5
295
    torch.testing.assert_close(output, ref_output, atol=atol, rtol=rtol)
296
297


298
299
300
301
302
303
304
305
306
def ref_multi_query_kv_attention(
    cu_seq_lens: List[int],
    query: torch.Tensor,
    key: torch.Tensor,
    value: torch.Tensor,
    scale: float,
    dtype: torch.dtype,
) -> torch.Tensor:
    num_seqs = len(cu_seq_lens) - 1
307
    ref_outputs: List[torch.Tensor] = []
308
309
310
311
312
313
314
315
316
    for i in range(num_seqs):
        start_idx = cu_seq_lens[i]
        end_idx = cu_seq_lens[i + 1]
        seq_len = end_idx - start_idx

        # Create attention mask.
        attn_mask = torch.triu(torch.ones(seq_len, seq_len, dtype=dtype),
                               diagonal=1)
        attn_mask = attn_mask * torch.finfo(dtype).min
317
        attn_mask = attn_mask.to(dtype=dtype)
318
319
320
321
322
323
324
325
326

        ref_output = ref_masked_attention(
            query[start_idx:end_idx],
            key[start_idx:end_idx],
            value[start_idx:end_idx],
            scale,
            attn_mask=attn_mask,
        )
        ref_outputs.append(ref_output)
327
328

    return torch.cat(ref_outputs, dim=0)
329
330


331
# TODO(woosuk): Add tests for USE_ALIBI=True.
332
333
334
335
336
@pytest.mark.parametrize("num_seqs", NUM_PREFILL_SEQS)
@pytest.mark.parametrize("num_heads", NUM_HEADS)
@pytest.mark.parametrize("head_size", HEAD_SIZES)
@pytest.mark.parametrize("dtype", DTYPES)
@pytest.mark.parametrize("seed", SEEDS)
337
@pytest.mark.parametrize("device", CUDA_DEVICES)
338
@torch.inference_mode()
339
def test_multi_query_kv_attention(
340
    num_seqs: int,
341
    num_heads: Tuple[int, int],
342
343
    head_size: int,
    dtype: torch.dtype,
344
    seed: int,
345
    device: str,
346
) -> None:
347
348
    random.seed(seed)
    torch.random.manual_seed(seed)
349
350
351
    if torch.cuda.is_available():
        torch.cuda.manual_seed(seed)
    torch.set_default_device(device)
352
353
354
355
356
    # MAX_SEQ_LEN sometimes causes OOM in the reference implementation.
    # As the xformers library is already tested with its own tests, we can use
    # a smaller MAX_SEQ_LEN here.
    max_len = min(MAX_SEQ_LEN, 4096)
    seq_lens = random.sample(range(1, max_len), num_seqs)
357
358
    num_tokens = sum(seq_lens)

359
    scale = float(1.0 / (head_size**0.5))
360
    num_query_heads, num_kv_heads = num_heads
361
    qkv = torch.empty(num_tokens,
362
                      num_query_heads + 2 * num_kv_heads,
363
                      head_size,
364
                      dtype=dtype)
365
366
367
368
369
370
371
372
373
    qkv.uniform_(-scale, scale)
    query, key, value = qkv.split(
        [num_query_heads, num_kv_heads, num_kv_heads], dim=1)

    num_queries_per_kv = num_query_heads // num_kv_heads
    if num_queries_per_kv > 1:
        # Handle MQA and GQA
        key = torch.repeat_interleave(key, num_queries_per_kv, dim=1)
        value = torch.repeat_interleave(value, num_queries_per_kv, dim=1)
374
375
376
377
378
379
380
381
    attn_bias = BlockDiagonalCausalMask.from_seqlens(seq_lens)
    output = xops.memory_efficient_attention_forward(
        query.unsqueeze(0),
        key.unsqueeze(0),
        value.unsqueeze(0),
        attn_bias=attn_bias,
        p=0.0,
        scale=scale,
Woosuk Kwon's avatar
Woosuk Kwon committed
382
    )
383
    output = output.squeeze(0)
384

385
386
387
    cu_seq_lens = [0]
    for seq_len in seq_lens:
        cu_seq_lens.append(cu_seq_lens[-1] + seq_len)
388
389
390
391
392
    ref_output = ref_multi_query_kv_attention(
        cu_seq_lens,
        query,
        key,
        value,
393
        scale,
394
395
        dtype,
    )
396
397
    atol = get_default_atol(output) if is_hip() else 1e-3
    rtol = get_default_rtol(output) if is_hip() else 1e-5
398
    torch.testing.assert_close(output, ref_output, atol=atol, rtol=rtol)