multi_step_worker.py 8.57 KB
Newer Older
1
import copy
2
import weakref
3
from typing import Dict, List, Tuple
4
5
6

import torch

7
from vllm.sequence import (ExecuteModelRequest, SamplerOutput, SequenceData,
8
                           SequenceGroupMetadata)
9
from vllm.spec_decode.interfaces import SpeculativeProposals
10
from vllm.spec_decode.proposer_worker_base import ProposerWorkerBase
11
from vllm.spec_decode.top1_proposer import Top1Proposer
12
from vllm.worker.worker import Worker
13
14


15
class MultiStepWorker(Worker, ProposerWorkerBase):
16
17
18
19
20
21
22
23
24
25
26
27
28
29
    """The MultiStepWorker is equivalent to a Worker except that it allows
    multiple forward passes in a single call, assuming the scheduler has
    allocated enough space to store the additional KV. This reduces overhead
    by invoking the scheduler less.

    The MultiStepWorker does not support cache swap operations, or beam search.
    Cache swap operations do not require large modifications. On the other hand,
    beam search requires memory allocations during sequence forks and thus
    requires more thought for MultiStepWorker support.
    """

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

30
        # Lazy initialization list.
31
        self._proposer: Top1Proposer
32

33
34
    def init_device(self):
        super().init_device()
35

36
        self._proposer = Top1Proposer(
37
            weakref.proxy(self),  # type: ignore[arg-type]
38
39
            self.device,
            self.vocab_size,
40
            max_proposal_len=self.max_model_len,
41
42
        )

43
44
45
46
    def set_include_gpu_probs_tensor(self):
        # Need include_gpu_probs_tensor for multi_step_worker
        self.model_runner.model.sampler.include_gpu_probs_tensor = True

47
    @torch.inference_mode()
48
    def sampler_output(
49
        self,
50
        execute_model_req: ExecuteModelRequest,
51
52
53
54
55
56
57
58
        sample_len: int,
    ) -> Tuple[List[SamplerOutput], bool]:
        """Run the model forward pass sample_len times. Returns the list of
        sampler output, one per model forward pass, along with indicator of
        whether torch tensor in sampler output need to be transposed in latter
        sampler_output_to_torch logic.

        For multi step worker, this indicator shall be True.
59
        """
60
        self._raise_if_unsupported(execute_model_req)
61
62
63
64

        # Shallow copy input data so modifications (such as appending tokens)
        # do not cause side-effects.
        copied_seq_group_metadata_list = self._shallow_copy_inputs(
65
66
67
            execute_model_req.seq_group_metadata_list)
        copied_execute_model_req = execute_model_req.clone(
            copied_seq_group_metadata_list)
68

69
        # Assert enough KV space for sample_len tokens per sequence.
70
71
        self._assert_enough_kv_space(execute_model_req.seq_group_metadata_list,
                                     sample_len)
72

73
        # Run model sample_len times.
74
        model_outputs: List[SamplerOutput] = []
75
        for _ in range(sample_len):
76
            model_output = super().execute_model(
77
                execute_model_req=copied_execute_model_req)
78
79
80
            assert (len(model_output) == 1
                    ), "composing multistep workers not supported"
            model_output = model_output[0]
81
82
83
84
85

            self._append_new_tokens(model_output,
                                    copied_seq_group_metadata_list)
            model_outputs.append(model_output)

86
        return model_outputs, True
87
88
89

    def get_spec_proposals(
        self,
90
        execute_model_req: ExecuteModelRequest,
91
92
93
94
95
    ) -> SpeculativeProposals:
        """Produce speculations given an input batch of sequences. The number of
        speculative tokens per sequence is determined by max_proposal_len.
        """

96
        return self._proposer.get_spec_proposals(execute_model_req)
97

98
    @staticmethod
99
    def _append_new_tokens(
100
101
            model_output: List[SamplerOutput],
            seq_group_metadata_list: List[SequenceGroupMetadata]) -> None:
102
103
104
105
106
107
108
109
110
111
112
113
114
115
116
117
118
        """Given model output from a single run, append the tokens to the
        sequences. This is normally done outside of the worker, but it is
        required if the worker is to perform multiple forward passes.
        """
        for seq_group_metadata, sequence_group_outputs in zip(
                seq_group_metadata_list, model_output):
            seq_group_metadata.is_prompt = False

            for seq_output in sequence_group_outputs.samples:
                # NOTE: Beam search is not supported, so we can assume that
                # parent_seq_id == seq_id.
                seq = seq_group_metadata.seq_data[seq_output.parent_seq_id]

                token_id = seq_output.output_token
                token_logprob = seq_output.logprobs[token_id]

                seq.append_token_id(token_id, token_logprob.logprob)
119
                seq.update_num_computed_tokens(1)
120

121
    @staticmethod
122
    def _shallow_copy_inputs(
123
        seq_group_metadata_list: List[SequenceGroupMetadata]
124
125
126
127
128
129
130
131
132
133
134
    ) -> List[SequenceGroupMetadata]:
        """Copy input data structures to remove side-effects when input data
        structures are shared with other modules.

        Helpful when the vLLM scheduler runs in the same process as the worker.
        The alternative is deep-copying (or other form of deep copy); this has
        performance downsides.
        """

        # Shallow-copy the list of SequenceGroupMetadata. This allows us to
        # append tokens and change is_prompt without external side-effects.
135
        new_seq_group_metadata_list: List[SequenceGroupMetadata] = []
136
137
138
139
140
141
142

        for old_seq_group_metadata in seq_group_metadata_list:
            # We must shallow-copy seq_group_metadata as is_prompt could change.
            seq_group_metadata = copy.copy(old_seq_group_metadata)
            new_seq_group_metadata_list.append(seq_group_metadata)

            # We must shallow-copy seq_data as we will append token ids
143
            new_seq_data: Dict[int, SequenceData] = {}
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
            for seq_id, old_seq_data in seq_group_metadata.seq_data.items():
                new_seq_data[seq_id] = copy.copy(old_seq_data)
                new_seq_data[
                    seq_id].output_token_ids = old_seq_data.output_token_ids[:]

            seq_group_metadata.seq_data = new_seq_data

        return new_seq_group_metadata_list

    def _assert_enough_kv_space(
            self, seq_group_metadata_list: List[SequenceGroupMetadata],
            num_steps: int) -> None:
        """Assert there are enough physical blocks per sequence to store the
        current KV plus additional KV from num_steps tokens.
        """
        assert self.model_runner.block_size is not None
        for seq_group_metadata in seq_group_metadata_list:
            # Only one seq_id is guaranteed because there is no beam search.
            seq_id = list(seq_group_metadata.seq_data.keys())[0]
            seq = seq_group_metadata.seq_data[seq_id]

            # After num_steps, the seq len will be the current seq len
            # plus one token per step.
            final_seq_len = seq.get_len() + num_steps

            # We will have final_seq_len - 1 KV because vLLM saves KV for a
            # token in the iteration after the token was generated.
            required_num_kv_slots = final_seq_len - 1

            # The allocated number of kv slots is the number of allocated blocks
            # times the number of slots of block.
            number_physical_blocks = len(
                seq_group_metadata.block_tables[seq_id])
            allocated_kv_slots = (number_physical_blocks *
                                  self.model_runner.block_size)

            if required_num_kv_slots > allocated_kv_slots:
                request_id = seq_group_metadata.request_id
                raise ValueError(
                    "The worker attempted to run "
                    f"{num_steps} times but found insufficient KV space for "
                    f"{request_id=} {seq_id=}. ({allocated_kv_slots=} "
                    f"{required_num_kv_slots=}).")

    def _raise_if_unsupported(
        self,
190
        execute_model_req: ExecuteModelRequest,
191
192
193
194
    ) -> None:
        """MultiStepWorker does not yet implement support for cache swap
        operations or beam search.
        """
195
196
197
198
199
        if any([
                execute_model_req.blocks_to_swap_in,
                execute_model_req.blocks_to_swap_out,
                execute_model_req.blocks_to_copy
        ]):
200
201
202
203
204
            raise NotImplementedError(
                "MultiStepWorker does not support cache operations")

        if any(
                len(seq_group_metadata.seq_data.keys()) != 1
205
206
                for seq_group_metadata in
                execute_model_req.seq_group_metadata_list):
207
208
            raise NotImplementedError(
                "MultiStepWorker does not support beam search.")