__init__.py 14.3 KB
Newer Older
1
# SPDX-License-Identifier: Apache-2.0
2
# SPDX-FileCopyrightText: Copyright contributors to the vLLM project
3
import multiprocessing
4
from concurrent.futures import Future, ThreadPoolExecutor
5
from typing import TYPE_CHECKING
6
7
8

from vllm.config import VllmConfig
from vllm.logger import init_logger
9
from vllm.reasoning import ReasoningParserManager
10
from vllm.tokenizers import init_tokenizer_from_config
11
from vllm.utils.import_utils import LazyLoader
12
from vllm.v1.structured_output.backend_guidance import GuidanceBackend
13
14
15
16
from vllm.v1.structured_output.backend_types import (
    StructuredOutputBackend,
    StructuredOutputGrammar,
)
17
from vllm.v1.structured_output.backend_xgrammar import XgrammarBackend
18
19
20
21

if TYPE_CHECKING:
    import numpy as np
    import numpy.typing as npt
22
    import torch
23

24
    from vllm.reasoning import ReasoningParser
25
    from vllm.v1.request import Request
26
27
else:
    torch = LazyLoader("torch", globals(), "torch")
28

29
30
31
    ReasoningParser = object
    Request = object

32
33
34
35
logger = init_logger(__name__)


class StructuredOutputManager:
36
    """Engine-level manager for structured output requests."""
37

38
    def __init__(self, vllm_config: VllmConfig):
39
40
        self.backend: StructuredOutputBackend | None = None
        self.reasoner: ReasoningParser | None = None
41
        self.vllm_config = vllm_config
42

43
44
45
46
47
48
49
50
51
52
        # When in external_launcher mode, async grammar compilation causes deadlocks
        # due to external_launcher mode having a scheduler for each TP rank.
        # Async grammar compilation causes the WAITING_FOR_FSM → WAITING transition to
        # happen at different times on different TP ranks,
        # breaking the determinism assumption that external_launcher relies on.
        self._use_async_grammar_compilation = (
            vllm_config.parallel_config.distributed_executor_backend
            != "external_launcher"
        )

53
        self._grammar_bitmask: torch.Tensor | None = None
54
        self._full_mask = torch.tensor(-1, dtype=torch.int32)
55

56
57
58
59
60
61
62
63
        max_batch_size = self.vllm_config.scheduler_config.max_num_seqs
        self.fill_bitmask_parallel_threshold = 128
        if self.fill_bitmask_parallel_threshold < max_batch_size:
            self.fill_bitmask_parallel_batch_size = 16
            # Use:
            # - at least 1 CPU
            # - at most half the number of CPUs or 8, whichever is less
            max_workers = max(1, min(multiprocessing.cpu_count() // 2, 8))
64
            self.executor_for_fillmask = ThreadPoolExecutor(max_workers=max_workers)
65

66
        if not vllm_config.renderer_config.skip_tokenizer_init:
67
68
69
70
71
72
73
            # The default max_workers if not specified is the number of
            # CPUs * 5, which is way too high since these tasks are CPU-bound,
            # not I/O bound. We also know we would never dominate CPU usage
            # with just grammar compilation, so we set it to half the number
            # of CPUs.
            max_workers = max(1, (multiprocessing.cpu_count() + 1) // 2)
            self.executor = ThreadPoolExecutor(max_workers=max_workers)
74
75
            self.tokenizer = init_tokenizer_from_config(vllm_config.renderer_config)
            reasoning_parser = vllm_config.structured_outputs_config.reasoning_parser
76
            reasoning_parser_plugin = (
77
                vllm_config.structured_outputs_config.reasoning_parser_plugin
78
79
80
81
            )
            if reasoning_parser_plugin and len(reasoning_parser_plugin) > 3:
                ReasoningParserManager.import_reasoning_parser(reasoning_parser_plugin)

82
            reasoning_parser = vllm_config.structured_outputs_config.reasoning_parser
83
            if reasoning_parser:
84
                reasoner_cls = ReasoningParserManager.get_reasoning_parser(
85
86
                    reasoning_parser
                )
87
                self.reasoner = reasoner_cls(tokenizer=self.tokenizer)
88

89
        self.enable_in_reasoning = (
90
            vllm_config.structured_outputs_config.enable_in_reasoning
91
92
        )

93
    def grammar_init(self, request: Request) -> None:
94
95
96
        if request.structured_output_request is None:
            return

97
        if TYPE_CHECKING:
98
99
100
101
            assert (
                request.sampling_params is not None
                and request.sampling_params.structured_outputs is not None
            )
102

103
104
105
106
        # Initialize the backend the first time it is needed.
        #
        # NOTE: We only support a single backend. We do NOT support different
        # backends on a per-request basis in V1 (for now, anyway...).
107
        # _backend is set in Processor._validate_structured_output
108
        if self.backend is None:
109
            assert request.sampling_params is not None
110
            backend = request.sampling_params.structured_outputs._backend
111
            vocab_size = self.vllm_config.model_config.get_vocab_size()
112
            if backend == "xgrammar":
113
114
115
116
117
                self.backend = XgrammarBackend(
                    self.vllm_config,
                    tokenizer=self.tokenizer,
                    vocab_size=vocab_size,
                )
118
            elif backend == "guidance":
119
120
121
122
123
                self.backend = GuidanceBackend(
                    self.vllm_config,
                    tokenizer=self.tokenizer,
                    vocab_size=vocab_size,
                )
124
            elif backend == "outlines":
125
                from vllm.v1.structured_output.backend_outlines import OutlinesBackend
126
127
128
129
130
131

                self.backend = OutlinesBackend(
                    self.vllm_config,
                    tokenizer=self.tokenizer,
                    vocab_size=vocab_size,
                )
132
133
            elif backend == "lm-format-enforcer":
                from vllm.v1.structured_output.backend_lm_format_enforcer import (  # noqa: E501
134
135
136
                    LMFormatEnforcerBackend,
                )

137
138
139
140
141
                self.backend = LMFormatEnforcerBackend(
                    self.vllm_config,
                    tokenizer=self.tokenizer,
                    vocab_size=vocab_size,
                )
142
            else:
143
                raise ValueError(f"Unsupported structured output backend: {backend}")
144

145
146
147
148
        if self._use_async_grammar_compilation:
            grammar = self.executor.submit(self._create_grammar, request)
        else:
            grammar = self._create_grammar(request)  # type: ignore[assignment]
149
        request.structured_output_request.grammar = grammar  # type: ignore[assignment]
150

151
    def _create_grammar(
152
153
154
        self,
        request: Request,
    ) -> StructuredOutputGrammar:
155
156
157
158
159
        key = request.structured_output_request.structured_output_key  # type: ignore[union-attr]

        # Note that the request was validated in the engine core client,
        # so at this point we know it is a supported type of request.
        #
160
161
        # TODO: we still need to handle xgrammar compilation failures,
        # though it should be unlikely as we test that up front as well.
162
163
        request_type, grammar_spec = key

164
165
        assert self.backend is not None
        return self.backend.compile_grammar(request_type, grammar_spec)
166

167
168
169
170
171
172
173
174
175
176
177
178
179
180
181
182
183
184
185
186
    def _fill_bitmasks(
        self,
        batch: list[tuple[StructuredOutputGrammar, int, bool]],
    ) -> None:
        assert self._grammar_bitmask is not None
        for grammar, index, apply_bitmask in batch:
            if apply_bitmask and not grammar.is_terminated():
                grammar.fill_bitmask(self._grammar_bitmask, index)
            else:
                # Note that for thinking support, we will need to
                # reset the relevant part of the bitmask for consequent
                # requests here.
                self._grammar_bitmask[index].fill_(self._full_mask)

    def _async_submit_fill_bitmask(
        self,
        batch: list[tuple[StructuredOutputGrammar, int, bool]],
    ) -> Future:
        return self.executor_for_fillmask.submit(self._fill_bitmasks, batch)

187
188
189
    def grammar_bitmask(
        self,
        requests: dict[str, Request],
190
        structured_output_request_ids: list[str],
191
        scheduled_spec_decode_tokens: dict[str, list[int]],
192
    ) -> "npt.NDArray[np.int32] | None":
193
194
195
196
        # Prepare the structured output bitmask for this batch.
        if not structured_output_request_ids:
            return None

197
198
        max_num_spec_tokens = 0
        if self.vllm_config.speculative_config is not None:
199
            max_num_spec_tokens = (
200
                self.vllm_config.speculative_config.num_speculative_tokens
201
            )
202

203
204
        if self._grammar_bitmask is None:
            assert self.backend is not None
205
206
207
208
209
            max_batch_size = self.vllm_config.scheduler_config.max_num_seqs

            # Allocate a bitmask for each token needing to be checked:
            # one for each speculative position, and one more for the
            # bonus token / non-speculative token.
210
211
212
            self._grammar_bitmask = self.backend.allocate_token_bitmask(
                max_batch_size * (1 + max_num_spec_tokens)
            )
213
214
215
216
217
218

        # Generate a batched bitmask for all structured output requests.
        # When speculative decoding is enabled, we need to include multiple
        # masks for each request, one for each possible bonus token position.
        # These are stored inline in the tensor and unpacked by the gpu runner.
        cumulative_index = 0
219

220
221
        # Optimized parallel filling of bitmasks for
        # non-spec, large-batch-size cases
222
        if (
223
            len(structured_output_request_ids) > self.fill_bitmask_parallel_threshold
224
225
            and max_num_spec_tokens == 0
        ):
226
227
            promises = []
            batch = []
228
            for req_id in structured_output_request_ids:
229
230
231
232
233
234
235
                request = requests[req_id]
                structured_output_request = request.structured_output_request
                if TYPE_CHECKING:
                    assert structured_output_request is not None
                    assert structured_output_request.grammar is not None

                apply_bitmask = self.should_fill_bitmask(request)
236
237
238
                batch.append(
                    (structured_output_request.grammar, cumulative_index, apply_bitmask)
                )
239
240
241
242
243
244
245
246
247
248
249
250
251
                if len(batch) == self.fill_bitmask_parallel_batch_size:
                    promises.append(self._async_submit_fill_bitmask(batch))
                    batch = []

                cumulative_index += 1
            if batch:
                promises.append(self._async_submit_fill_bitmask(batch))

            # Wait for all bitmask filling tasks to complete.
            for promise in promises:
                promise.result()
        else:
            # Fallback to serial filling of bitmasks for small-batch-size cases
252
            for req_id in structured_output_request_ids:
253
254
255
256
257
258
259
260
261
262
263
                request = requests[req_id]
                structured_output_request = request.structured_output_request

                if TYPE_CHECKING:
                    assert structured_output_request is not None
                    assert structured_output_request.grammar is not None
                apply_bitmask = self.should_fill_bitmask(request)

                state_advancements = 0
                req_tokens = scheduled_spec_decode_tokens.get(req_id, [])
                for i, token in enumerate(req_tokens + [None]):
264
265
266
267
268
269
270
271
272
273
274
275
276
277
278
                    self._fill_bitmasks(
                        [
                            (
                                structured_output_request.grammar,
                                cumulative_index,
                                apply_bitmask,
                            )
                        ]
                    )

                    if (
                        apply_bitmask
                        and token is not None
                        and not structured_output_request.grammar.is_terminated()
                    ):
279
                        accepted = structured_output_request.grammar.accept_tokens(
280
281
                            req_id, [token]
                        )
282
                        assert accepted, (token, req_id, scheduled_spec_decode_tokens)
283
                        state_advancements += 1
284
285
                    cumulative_index += 1
                if state_advancements > 0:
286
                    structured_output_request.grammar.rollback(state_advancements)
287

288
        bitmask_tensor = self._grammar_bitmask
289
290
        if cumulative_index < bitmask_tensor.shape[0]:
            bitmask_tensor = bitmask_tensor[:cumulative_index]
291
292
293
294
295

        # After finishing with the xgrammar operations, we convert to
        # np.ndarray, because that is much more efficient for serialization
        # and deserialization when sending this to the GPU workers.
        return bitmask_tensor.numpy()
296

297
    def should_fill_bitmask(self, request: Request) -> bool:
298
299
300
301
        # NOTE (Hanchen) if enable_in_reasoning is True, it means that
        # the model needs to be constrained in reasoning. So we should always
        # enable the bitmask filling.

302
        if self.reasoner is not None:
303
304
            if self.enable_in_reasoning:
                return True
305
306
            assert request.structured_output_request is not None
            if request.structured_output_request.reasoning_ended is None:
307
                request.structured_output_request.reasoning_ended = (
308
                    self.reasoner.is_reasoning_end(request.prompt_token_ids)
309
                )
310
311
312
            return request.structured_output_request.reasoning_ended
        return True

313
314
315
316
317
318
319
320
321
322
    def should_advance(self, request: Request) -> bool:
        if not request.use_structured_output:
            return False

        # To determine whether we can advance the FSM.
        # Supports thinking usage where we skip the reasoning components.
        if TYPE_CHECKING:
            assert request.structured_output_request is not None
            assert request.structured_output_request.grammar is not None
        # by default, we should always advance
323
        # for cases that don't use thinking mode.
324
325
        if self.reasoner is None:
            return True
326

327
328
329
330
        # if the model needs structured in reasoning, we should advance
        if self.enable_in_reasoning:
            return True

331
332
333
        structured_req = request.structured_output_request
        if structured_req.reasoning_ended:
            return True
334

335
336
337
338
339
        # Check if reasoning ends in *this* step
        if self.reasoner.is_reasoning_end(request.all_token_ids):
            # Reasoning just ended, so we shouldn't advance til
            # next pass
            structured_req.reasoning_ended = True
340

341
        return False
342

343
344
345
    def clear_backend(self) -> None:
        if self.backend is not None:
            self.backend.destroy()