test_spec_decode.py 5.71 KB
Newer Older
1
# SPDX-License-Identifier: Apache-2.0
2
# SPDX-FileCopyrightText: Copyright contributors to the vLLM project
3
4
from __future__ import annotations

5
import random
6
from typing import Any
7

8
import pytest
zhiweiz's avatar
zhiweiz committed
9
import torch
10
11

from vllm import LLM, SamplingParams
zhiweiz's avatar
zhiweiz committed
12
from vllm.distributed import cleanup_dist_env_and_memory
13
14
15
16


@pytest.fixture
def test_prompts():
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
    prompt_types = ["repeat", "sentence"]
    num_prompts = 100
    prompts = []

    random.seed(0)
    random_prompt_type_choices = random.choices(prompt_types, k=num_prompts)

    # Generate a mixed batch of prompts, some of which can be easily
    # predicted by n-gram matching and some which likely cannot.
    for kind in random_prompt_type_choices:
        word_choices = ["test", "temp", "hello", "where"]
        word = random.choice(word_choices)
        if kind == "repeat":
            prompt = f"""
            please repeat the word '{word}' 10 times.
            give no other output than the word at least ten times in a row,
            in lowercase with spaces between each word and without quotes.
            """
        elif kind == "sentence":
            prompt = f"""
            please give a ten-word sentence that
            uses the word {word} at least once.
            give no other output than that simple sentence without quotes.
            """
        else:
            raise ValueError(f"Unknown prompt type: {kind}")
        prompts.append([{"role": "user", "content": prompt}])

    return prompts
46
47
48
49


@pytest.fixture
def sampling_config():
50
    return SamplingParams(temperature=0, max_tokens=10, ignore_eos=False)
51
52
53
54


@pytest.fixture
def model_name():
55
    return "meta-llama/Llama-3.1-8B-Instruct"
56
57


58
59
60
61
62
63
def test_ngram_correctness(
    monkeypatch: pytest.MonkeyPatch,
    test_prompts: list[list[dict[str, Any]]],
    sampling_config: SamplingParams,
    model_name: str,
):
64
65
66
67
68
69
70
    '''
    Compare the outputs of a original LLM and a speculative LLM
    should be the same when using ngram speculative decoding.
    '''
    with monkeypatch.context() as m:
        m.setenv("VLLM_USE_V1", "1")

71
72
        ref_llm = LLM(model=model_name, max_model_len=1024)
        ref_outputs = ref_llm.chat(test_prompts, sampling_config)
73
        del ref_llm
zhiweiz's avatar
zhiweiz committed
74
75
        torch.cuda.empty_cache()
        cleanup_dist_env_and_memory()
76

77
78
79
80
81
82
83
84
85
86
        spec_llm = LLM(
            model=model_name,
            speculative_config={
                "method": "ngram",
                "prompt_lookup_max": 5,
                "prompt_lookup_min": 3,
                "num_speculative_tokens": 3,
            },
            max_model_len=1024,
        )
87
88
89
        spec_outputs = spec_llm.chat(test_prompts, sampling_config)
        matches = 0
        misses = 0
90
        for ref_output, spec_output in zip(ref_outputs, spec_outputs):
91
92
93
94
95
96
97
98
99
100
            if ref_output.outputs[0].text == spec_output.outputs[0].text:
                matches += 1
            else:
                misses += 1
                print(f"ref_output: {ref_output.outputs[0].text}")
                print(f"spec_output: {spec_output.outputs[0].text}")

        # Heuristic: expect at least 70% of the prompts to match exactly
        # Upon failure, inspect the outputs to check for inaccuracy.
        assert matches > int(0.7 * len(ref_outputs))
101
        del spec_llm
zhiweiz's avatar
zhiweiz committed
102
103
104
105
106
107
108
109
110
111
112
113
114
115
116
        torch.cuda.empty_cache()
        cleanup_dist_env_and_memory()


@pytest.mark.parametrize("model_setup", [
    ("eagle", "meta-llama/Llama-3.1-8B-Instruct",
     "yuhuili/EAGLE-LLaMA3.1-Instruct-8B", 1),
    ("eagle3", "meta-llama/Llama-3.1-8B-Instruct",
     "yuhuili/EAGLE3-LLaMA3.1-Instruct-8B", 1),
    pytest.param(
        ("eagle", "meta-llama/Llama-4-Scout-17B-16E-Instruct",
         "morgendave/EAGLE-Llama-4-Scout-17B-16E-Instruct", 4),
        marks=pytest.mark.skip(reason="Skipping due to CI OOM issues")),
],
                         ids=["llama3_eagle", "llama3_eagle3", "llama4_eagle"])
117
118
119
120
def test_eagle_correctness(
    monkeypatch: pytest.MonkeyPatch,
    test_prompts: list[list[dict[str, Any]]],
    sampling_config: SamplingParams,
zhiweiz's avatar
zhiweiz committed
121
    model_setup: tuple[str, str, str, int],
122
123
124
125
):
    '''
    Compare the outputs of a original LLM and a speculative LLM
    should be the same when using eagle speculative decoding.
zhiweiz's avatar
zhiweiz committed
126
    model_setup: (method, model_name, eagle_model_name, tp_size)
127
128
129
    '''
    with monkeypatch.context() as m:
        m.setenv("VLLM_USE_V1", "1")
zhiweiz's avatar
zhiweiz committed
130
        method, model_name, spec_model_name, tp_size = model_setup
131

zhiweiz's avatar
zhiweiz committed
132
133
134
        ref_llm = LLM(model=model_name,
                      max_model_len=2048,
                      tensor_parallel_size=tp_size)
135
136
        ref_outputs = ref_llm.chat(test_prompts, sampling_config)
        del ref_llm
zhiweiz's avatar
zhiweiz committed
137
138
        torch.cuda.empty_cache()
        cleanup_dist_env_and_memory()
139
140
141

        spec_llm = LLM(
            model=model_name,
142
            trust_remote_code=True,
zhiweiz's avatar
zhiweiz committed
143
            tensor_parallel_size=tp_size,
144
            speculative_config={
zhiweiz's avatar
zhiweiz committed
145
                "method": method,
146
                "model": spec_model_name,
147
                "num_speculative_tokens": 3,
148
                "max_model_len": 2048,
149
            },
150
            max_model_len=2048,
151
152
153
154
155
156
157
158
159
160
161
162
        )
        spec_outputs = spec_llm.chat(test_prompts, sampling_config)
        matches = 0
        misses = 0
        for ref_output, spec_output in zip(ref_outputs, spec_outputs):
            if ref_output.outputs[0].text == spec_output.outputs[0].text:
                matches += 1
            else:
                misses += 1
                print(f"ref_output: {ref_output.outputs[0].text}")
                print(f"spec_output: {spec_output.outputs[0].text}")

163
        # Heuristic: expect at least 66% of the prompts to match exactly
164
        # Upon failure, inspect the outputs to check for inaccuracy.
165
        assert matches > int(0.66 * len(ref_outputs))
166
        del spec_llm
zhiweiz's avatar
zhiweiz committed
167
168
        torch.cuda.empty_cache()
        cleanup_dist_env_and_memory()