test_fusion_attn.py 16.3 KB
Newer Older
1
# SPDX-License-Identifier: Apache-2.0
2
# SPDX-FileCopyrightText: Copyright contributors to the vLLM project
3
import copy
4
5
6
7

import pytest
import torch._dynamo

8
from tests.compile.backend import LazyInitPass, TestBackend
9
from tests.v1.attention.utils import BatchSpec, create_common_attn_metadata
10
from vllm._custom_ops import cutlass_scaled_fp4_mm, scaled_fp4_quant
11
from vllm.attention import Attention, AttentionMetadata
12
from vllm.attention.backends.registry import _Backend
13
from vllm.attention.selector import global_force_attn_backend_context_manager
14
from vllm.compilation.fusion import QUANT_OPS
15
16
17
from vllm.compilation.fusion_attn import ATTN_OP, AttnFusionPass
from vllm.compilation.fx_utils import find_op_nodes
from vllm.compilation.noop_elimination import NoOpEliminationPass
18
from vllm.compilation.post_cleanup import PostCleanupPass
19
20
21
22
23
24
25
26
27
28
from vllm.config import (
    CacheConfig,
    CompilationConfig,
    CompilationLevel,
    ModelConfig,
    PassConfig,
    SchedulerConfig,
    VllmConfig,
    set_current_vllm_config,
)
29
30
from vllm.forward_context import get_forward_context, set_forward_context
from vllm.model_executor.layers.quantization.utils.quant_utils import (
31
32
33
34
    kFp8StaticTensorSym,
    kNvfp4Quant,
)
from vllm.model_executor.layers.quantization.utils.w8a8_utils import Fp8LinearOp
35
from vllm.platforms import current_platform
36
from vllm.utils import is_torch_equal_or_newer
37
38
39
from vllm.v1.kv_cache_interface import AttentionSpec

FP8_DTYPE = current_platform.fp8_dtype()
40
FP4_DTYPE = torch.uint8
41
42

# globals needed for string-import custom Dynamo backend field
43
44
backend: TestBackend | None = None
backend_unfused: TestBackend | None = None
45
46


47
48
class AttentionQuantPatternModel(torch.nn.Module):
    """Base model for AttentionQuantPattern fusion."""
49

50
51
52
53
54
55
56
57
58
59
    def __init__(
        self,
        num_qo_heads: int,
        num_kv_heads: int,
        head_size: int,
        kv_cache_dtype: torch.dtype,
        device: torch.device,
        vllm_config: VllmConfig,
        **kwargs,
    ):
60
61
62
63
64
65
66
67
68
69
70
71
72
73
74
75
        super().__init__()
        self.num_qo_heads = num_qo_heads
        self.num_kv_heads = num_kv_heads
        self.head_size = head_size
        self.kv_cache_dtype = kv_cache_dtype
        self.device = device
        self.vllm_config = vllm_config

        self.attn = Attention(
            num_heads=self.num_qo_heads,
            head_size=self.head_size,
            scale=1.0 / (self.head_size**0.5),
            num_kv_heads=self.num_kv_heads,
            cache_config=vllm_config.cache_config,
            prefix="model.layers.0.self_attn.attn",
        )
76
77
        self.attn._k_scale = self.attn._k_scale.to(device)
        self.attn._v_scale = self.attn._v_scale.to(device)
78
79
80
81
82
83
84
85
86
87
88
89
90
91
92
93

        self.block_size = 16

        # Initialize attn MetadataBuilder
        self.builder = self.attn.attn_backend.get_builder_cls()(
            kv_cache_spec=AttentionSpec(
                block_size=self.block_size,
                num_kv_heads=self.num_kv_heads,
                head_size=self.head_size,
                dtype=self.kv_cache_dtype,
            ),
            layer_names=[self.attn.layer_name],
            vllm_config=self.vllm_config,
            device=self.device,
        )

94
    def build_attn_metadata(self, batch_size: int) -> AttentionMetadata:
95
96
97
        """Initialize attention metadata."""

        # Create common attn metadata
98
        batch_spec = BatchSpec(seq_lens=[1] * batch_size, query_lens=[1] * batch_size)
99
        common_attn_metadata = create_common_attn_metadata(
100
101
            batch_spec, self.block_size, self.device, arange_block_indices=True
        )
102

103
        max_blocks = (max(batch_spec.seq_lens) + self.block_size - 1) // self.block_size
104
        num_blocks = batch_size * max_blocks
105
        backend = self.attn.backend
106

107
108
        # Create dummy KV cache for the selected backend
        if backend == _Backend.ROCM_ATTN:
109
            # k/v as 1st dimention
110
111
112
113
114
115
116
117
118
119
120
121
122
123
124
125
126
127
128
129
130
131
132
            # HND: [num_blocks, num_kv_heads, block_size, head_size]
            kv_cache = torch.zeros(
                2,
                num_blocks,
                self.num_kv_heads,
                self.block_size,
                self.head_size,
                dtype=self.kv_cache_dtype,
                device=self.device,
            )
        elif backend == _Backend.ROCM_AITER_UNIFIED_ATTN:
            # k/v as 1st dimention
            # NHD: [num_blocks, block_size, num_kv_heads, head_size]
            kv_cache = torch.zeros(
                2,
                num_blocks,
                self.block_size,
                self.num_kv_heads,
                self.head_size,
                dtype=self.kv_cache_dtype,
                device=self.device,
            )
        elif backend == _Backend.TRITON_ATTN:
133
            # k/v as 2nd dimention
134
135
136
137
138
139
140
141
142
143
144
145
146
147
148
149
150
151
152
153
154
155
            # NHD: [num_blocks, block_size, num_kv_heads, head_size]
            kv_cache = torch.zeros(
                num_blocks,
                2,
                self.num_kv_heads,
                self.block_size,
                self.head_size,
                dtype=self.kv_cache_dtype,
                device=self.device,
            )
        elif backend == _Backend.FLASHINFER:
            kv_cache = torch.zeros(
                num_blocks,
                2,
                self.num_kv_heads,
                self.block_size,
                self.head_size,
                dtype=self.kv_cache_dtype,
                device=self.device,
            ).permute(0, 1, 3, 2, 4)
        else:
            raise ValueError(f"Unsupported backend: {backend}")
156
157
158
159
        self.attn.kv_cache = [kv_cache]

        # Build attn metadata
        self.attn_metadata = self.builder.build(
160
161
            common_prefix_len=0, common_attn_metadata=common_attn_metadata
        )
162
163
164

        return self.attn_metadata

165
166
167
168
169
170
171
172
173
174
175

class TestAttentionFp8StaticQuantPatternModel(AttentionQuantPatternModel):
    """Test model for AttentionFp8StaticQuantPattern fusion."""

    quant_key = kFp8StaticTensorSym

    def __init__(self, *args, **kwargs):
        super().__init__(*args, **kwargs)

        self.fp8_linear = Fp8LinearOp(
            act_quant_static=self.quant_key.scale.static,
176
177
            act_quant_group_shape=self.quant_key.scale.group_shape,
        )
178
179
180

        hidden_size = self.num_qo_heads * self.head_size
        self.w = kwargs.get(
181
182
183
184
185
186
187
188
189
            "w",
            {
                "weight": torch.randn(hidden_size, hidden_size)
                .to(dtype=FP8_DTYPE, device=self.device)
                .t(),
                "wscale": torch.tensor([1.0], dtype=torch.float32, device=self.device),
                "scale": torch.tensor([1.0], dtype=torch.float32, device=self.device),
            },
        )
190
191

    def forward(self, q: torch.Tensor, k: torch.Tensor, v: torch.Tensor):
192
193
        """Forward pass that creates the pattern to be fused."""
        attn_output = self.attn(q, k, v)
194
195
196
197
198
199
        return self.fp8_linear.apply(
            input=attn_output,
            weight=self.w["weight"],
            weight_scale=self.w["wscale"],
            input_scale=self.w["scale"],
        )
200
201
202
203
204
205
206
207
208
209
210
211


class TestAttentionNvfp4QuantPatternModel(AttentionQuantPatternModel):
    """Test model for AttentionNvfp4QuantPattern fusion."""

    quant_key = kNvfp4Quant

    def __init__(self, *args, **kwargs):
        super().__init__(*args, **kwargs)

        hidden_size = self.num_qo_heads * self.head_size
        self.w = kwargs.get(
212
213
214
215
216
217
218
219
220
221
222
223
224
225
226
            "w",
            {
                "weight": torch.randint(
                    256,
                    (hidden_size, hidden_size // 2),
                    dtype=FP4_DTYPE,
                    device=self.device,
                ),
                "wscale_swizzled": torch.randn(hidden_size, hidden_size // 16).to(
                    dtype=FP8_DTYPE, device=self.device
                ),
                "wscale": torch.tensor([500], dtype=torch.float32, device=self.device),
                "scale": torch.tensor([0.002], dtype=torch.float32, device=self.device),
            },
        )
227
228
229
230
231

    def forward(self, q: torch.Tensor, k: torch.Tensor, v: torch.Tensor):
        """Forward pass that creates the pattern to be fused."""
        attn_output = self.attn(q, k, v)
        quant_output, output_block_scale = scaled_fp4_quant(
232
233
234
235
236
237
238
239
240
241
            attn_output, 1 / self.w["scale"]
        )
        return cutlass_scaled_fp4_mm(
            a=quant_output,
            b=self.w["weight"],
            block_scale_a=output_block_scale,
            block_scale_b=self.w["wscale_swizzled"],
            alpha=self.w["scale"] * self.w["wscale"],
            out_dtype=attn_output.dtype,
        )
242
243


244
if current_platform.is_cuda():
245
246
247
248
249
250
251
252
253
254
    MODELS = [
        (
            "nvidia/Llama-4-Scout-17B-16E-Instruct-FP8",
            TestAttentionFp8StaticQuantPatternModel,
        ),
        (
            "nvidia/Llama-4-Scout-17B-16E-Instruct-FP4",
            TestAttentionNvfp4QuantPatternModel,
        ),
    ]
255
256
    HEADS = [(64, 8), (40, 8)]
elif current_platform.is_rocm():
257
258
259
    MODELS = [
        ("amd/Llama-3.1-8B-Instruct-FP8-KV", TestAttentionFp8StaticQuantPatternModel)
    ]
260
261
262
263
264
265
266
    HEADS = [(32, 8), (40, 8)]
else:
    MODELS = []
    HEADS = []


@pytest.mark.parametrize("num_qo_heads, num_kv_heads", HEADS)
267
@pytest.mark.parametrize("head_size", [128])
268
269
270
@pytest.mark.parametrize(
    "batch_size", [7, 256, 533] if current_platform.is_cuda() else [8]
)
271
272
273
@pytest.mark.parametrize("dtype", [torch.bfloat16, torch.float16])
@pytest.mark.parametrize("model_name, model_class", MODELS)
@pytest.mark.parametrize(
274
    "backend",
275
276
277
    [_Backend.FLASHINFER]
    if current_platform.is_cuda()
    else [_Backend.ROCM_AITER_UNIFIED_ATTN, _Backend.ROCM_ATTN, _Backend.TRITON_ATTN],
278
)
279
280
281
# TODO(boyuan): test inductor graph partition on rocm
@pytest.mark.parametrize(
    "use_inductor_graph_partition",
282
283
284
285
286
    [False] if current_platform.is_rocm() else [False, True],
)
@pytest.mark.skipif(
    not current_platform.is_cuda_alike(), reason="Only test ROCm or CUDA"
)
287
@pytest.mark.skipif(not current_platform.supports_fp8(), reason="Need FP8")
288
289
290
291
292
293
294
295
296
297
298
299
300
301
302
303
304
305
306
307
@pytest.mark.skipif(
    current_platform.is_cuda() and not current_platform.is_device_capability((10, 0)),
    reason="On CUDA only test on SM100(Blackwell)",
)
@pytest.mark.skipif(
    not current_platform.is_cuda_alike(), reason="Only test ROCm or CUDA"
)
def test_attention_quant_pattern(
    num_qo_heads: int,
    num_kv_heads: int,
    head_size: int,
    batch_size: int,
    dtype: torch.dtype,
    model_name: str,
    model_class: type[AttentionQuantPatternModel],
    backend: _Backend,
    use_inductor_graph_partition: bool,
    dist_init,
    caplog_vllm,
):
308
309
    """Test AttentionStaticQuantPattern fusion pass"""

310
311
    if use_inductor_graph_partition and not is_torch_equal_or_newer("2.9.0.dev"):
        pytest.skip("inductor graph partition is only available in PyTorch 2.9+")
312

313
314
315
316
317
318
319
    device = torch.device("cuda:0")
    torch.manual_seed(42)

    vllm_config = VllmConfig(
        model_config=ModelConfig(
            model=model_name,
            max_model_len=2048,
320
            dtype=dtype,
321
322
323
324
325
        ),
        scheduler_config=SchedulerConfig(max_num_seqs=1024),
        compilation_config=CompilationConfig(
            level=CompilationLevel.PIECEWISE,
            custom_ops=["+quant_fp8"],
326
            use_inductor_graph_partition=use_inductor_graph_partition,
327
        ),
328
329
        cache_config=CacheConfig(cache_dtype="fp8"),
    )
330
331

    # Create test inputs
332
333
334
    q = torch.randn(batch_size, num_qo_heads * head_size, dtype=dtype, device=device)
    k = torch.randn(batch_size, num_kv_heads * head_size, dtype=dtype, device=device)
    v = torch.randn(batch_size, num_kv_heads * head_size, dtype=dtype, device=device)
335
336
337
338
339
340
341
342

    # Mark first dimension as dynamic for realistic testing
    torch._dynamo.mark_dynamic(q, 0)
    torch._dynamo.mark_dynamic(k, 0)
    torch._dynamo.mark_dynamic(v, 0)

    # Run model directly without compilation and fusion
    vllm_config_unfused = copy.deepcopy(vllm_config)
343
344
345
346
347
348
349
350
351
352
353
354
355
    with (
        set_current_vllm_config(vllm_config_unfused),
        set_forward_context(attn_metadata=None, vllm_config=vllm_config_unfused),
        global_force_attn_backend_context_manager(backend),
    ):
        model_unfused = model_class(
            num_qo_heads=num_qo_heads,
            num_kv_heads=num_kv_heads,
            head_size=head_size,
            kv_cache_dtype=FP8_DTYPE,
            device=device,
            vllm_config=vllm_config_unfused,
        )
356
357
358
        model_unfused = model_unfused.to(device)

        forward_ctx = get_forward_context()
359
        forward_ctx.attn_metadata = model_unfused.build_attn_metadata(batch_size)
360
361

        # Run model directly without compilation and fusion
362
        result_unfused = model_unfused(q, k, v)
363
364
365

    # Run model with attn fusion enabled
    vllm_config.compilation_config.pass_config = PassConfig(
366
367
368
369
370
371
372
373
374
375
376
377
378
379
380
381
        enable_attn_fusion=True, enable_noop=True
    )
    with (
        set_current_vllm_config(vllm_config),
        set_forward_context(attn_metadata=None, vllm_config=vllm_config),
        global_force_attn_backend_context_manager(backend),
    ):
        model_fused = model_class(
            num_qo_heads=num_qo_heads,
            num_kv_heads=num_kv_heads,
            head_size=head_size,
            kv_cache_dtype=FP8_DTYPE,
            device=device,
            vllm_config=vllm_config,
            w=model_unfused.w,
        )
382
383
384
        model_fused = model_fused.to(device)

        forward_ctx = get_forward_context()
385
        forward_ctx.attn_metadata = model_fused.build_attn_metadata(batch_size)
386
387
388

        # Create test backend with fusion passes enabled
        noop_pass = NoOpEliminationPass(vllm_config)
389
390
391
392
        attn_pass = LazyInitPass(AttnFusionPass, vllm_config)
        cleanup_pass = PostCleanupPass(vllm_config)

        test_backend = TestBackend(noop_pass, attn_pass, cleanup_pass)
393
394

        # Compile model with fusion enabled
395
396
397
        model_compiled = torch.compile(
            model_fused, backend=test_backend, fullgraph=True
        )
398
        assert model_compiled.attn._o_scale_float is None
399

400
        result_fused_1 = model_compiled(q, k, v)
401

402
403
404
405
406
407
408
        if backend == _Backend.FLASHINFER:
            # With the Flashinfer backend after the 1st round of the forward
            # pass, output quant scale should be loaded into the attn layer's
            # _o_scale_float, the 2nd round should reuse the loaded
            # _o_scale_float
            assert model_compiled.attn._o_scale_float is not None
            result_fused_2 = model_compiled(q, k, v)
409

410
411
            assert model_compiled.attn._o_scale_float is not None

412
413
414
            torch.testing.assert_close(
                result_unfused, result_fused_2, atol=1e-2, rtol=1e-2
            )
415
416

    # Check attn fusion support
417
    quant_key = model_class.quant_key
418
    attn_fusion_supported = [
419
420
        layer.impl.fused_output_quant_supported(quant_key)
        for key, layer in vllm_config.compilation_config.static_forward_context.items()
421
422
423
    ]
    if any(attn_fusion_supported):
        # Check quantization ops in the graph before and after fusion
424
        test_backend.check_before_ops([QUANT_OPS[quant_key]], fully_replaced=True)
425

426
427
428
    # access the underlying `AttnFusionPass` on the `LazyInitPass`
    assert attn_pass.pass_.matched_count == sum(attn_fusion_supported)

429
430
    # Check attention ops in the graph before and after fusion
    attn_nodes_pre = list(find_op_nodes(ATTN_OP, test_backend.graph_pre_pass))
431
    attn_nodes_post = list(find_op_nodes(ATTN_OP, test_backend.graph_post_pass))
432
433

    assert len(attn_nodes_pre) > 0, "Should have attention nodes before fusion"
434
    assert len(attn_nodes_pre) == len(attn_nodes_post), (
435
        "Should have same number of attention nodes before and after fusion"
436
437
    )
    assert attn_nodes_pre[0].kwargs.get("output_scale") is None, (
438
        "Attention should not have output_scale before fusion"
439
440
    )
    assert attn_nodes_post[0].kwargs.get("output_scale") is not None, (
441
        "Attention should have output_scale after fusion"
442
    )
443

444
    assert attn_nodes_pre[0].kwargs.get("output_block_scale") is None, (
445
        "Attention should not have output_block_scale before fusion"
446
    )
447
    if quant_key.dtype == FP8_DTYPE:
448
        assert attn_nodes_post[0].kwargs.get("output_block_scale") is None, (
449
            "Attention should not have output_block_scale after FP8 fusion"
450
        )
451
    elif quant_key.dtype == FP4_DTYPE:
452
453
        assert attn_nodes_post[0].kwargs.get("output_block_scale") is not None, (
            "Attention should have output_block_scale after FP4 fusion"
454
        )
455

456
    # Check that results are close
457
    torch.testing.assert_close(result_unfused, result_fused_1, atol=1e-2, rtol=1e-2)