Skip to content

Commit

Permalink
feature(xyy):add HPT model and test_hpt
Browse files Browse the repository at this point in the history
  • Loading branch information
luodi-7 committed Dec 4, 2024
1 parent 9608131 commit 95a3edc
Show file tree
Hide file tree
Showing 4 changed files with 29 additions and 16 deletions.
12 changes: 8 additions & 4 deletions ding/model/template/hpt.py
Original file line number Diff line number Diff line change
Expand Up @@ -97,7 +97,8 @@ def init_cross_attn(self):
"""Initialize cross-attention module and learnable tokens."""
token_num = 16
self.tokens = nn.Parameter(torch.randn(1, token_num, 128) * INIT_CONST)
self.cross_attention = CrossAttention(128, heads=8, dim_head=64, dropout=0.1)
self.cross_attention = CrossAttention(
128, heads=8, dim_head=64, dropout=0.1)

def compute_latent(self, x: torch.Tensor) -> torch.Tensor:
"""
Expand All @@ -112,10 +113,12 @@ def compute_latent(self, x: torch.Tensor) -> torch.Tensor:
"""
# Using the Feature Extractor
stem_feat = self.feature_extractor(x)
stem_feat = stem_feat.reshape(stem_feat.shape[0], -1, stem_feat.shape[-1]) # (B, N, 128)
stem_feat = stem_feat.reshape(
stem_feat.shape[0], -1, stem_feat.shape[-1]) # (B, N, 128)
# Calculating latent tokens using CrossAttention
stem_tokens = self.tokens.repeat(len(stem_feat), 1, 1) # (B, 16, 128)
stem_tokens = self.cross_attention(stem_tokens, stem_feat) # (B, 16, 128)
stem_tokens = self.cross_attention(
stem_tokens, stem_feat) # (B, 16, 128)
return stem_tokens

def forward(self, x: torch.Tensor) -> torch.Tensor:
Expand Down Expand Up @@ -195,7 +198,8 @@ def forward(self, x: torch.Tensor, context: torch.Tensor, mask: Optional[torch.T
h = self.heads
q = self.to_q(x)
k, v = self.to_kv(context).chunk(2, dim=-1)
q, k, v = map(lambda t: rearrange(t, "b n (h d) -> (b h) n d", h=h), (q, k, v))
q, k, v = map(lambda t: rearrange(
t, "b n (h d) -> (b h) n d", h=h), (q, k, v))
sim = torch.einsum("b i d, b j d -> b i j", q, k) * self.scale

if mask is not None:
Expand Down
6 changes: 3 additions & 3 deletions ding/model/template/tests/test_hpt.py
Original file line number Diff line number Diff line change
Expand Up @@ -4,9 +4,9 @@
from ding.model.template.hpt import HPT
from ding.torch_utils import is_differentiable

T, B = 3, 4
obs_shape = [4, (8, ), (4, 64, 64)] # Example observation shapes
act_shape = [3, (6, ), [2, 3, 6]] # Example action shapes
T, B = 4
obs_shape = [4, (8, )]
act_shape = [3, (6, )]
args = list(product(*[obs_shape, act_shape]))


Expand Down
12 changes: 8 additions & 4 deletions dizoo/box2d/lunarlander/entry/lunarlander_dqn_example.py
Original file line number Diff line number Diff line change
Expand Up @@ -19,16 +19,19 @@

def main():
logging.getLogger().setLevel(logging.INFO)
cfg = compile_config(main_config, create_cfg=create_config, auto=True, save_cfg=task.router.node_id == 0)
cfg = compile_config(main_config, create_cfg=create_config,
auto=True, save_cfg=task.router.node_id == 0)
ding_init(cfg)

with task.start(async_mode=False, ctx=OnlineRLContext()):
collector_env = SubprocessEnvManagerV2(
env_fn=[lambda: DingEnvWrapper(gym.make("LunarLander-v2")) for _ in range(cfg.env.collector_env_num)],
env_fn=[lambda: DingEnvWrapper(
gym.make("LunarLander-v2")) for _ in range(cfg.env.collector_env_num)],
cfg=cfg.env.manager
)
evaluator_env = SubprocessEnvManagerV2(
env_fn=[lambda: DingEnvWrapper(gym.make("LunarLander-v2")) for _ in range(cfg.env.evaluator_env_num)],
env_fn=[lambda: DingEnvWrapper(
gym.make("LunarLander-v2")) for _ in range(cfg.env.evaluator_env_num)],
cfg=cfg.env.manager
)

Expand All @@ -38,7 +41,8 @@ def main():
device = torch.device("cuda" if torch.cuda.is_available() else "cpu")
model = DQN(**cfg.policy.model).to(device)

buffer_ = DequeBuffer(size=cfg.policy.other.replay_buffer.replay_buffer_size)
buffer_ = DequeBuffer(
size=cfg.policy.other.replay_buffer.replay_buffer_size)

# Pass the model into Policy
policy = DQNPolicy(cfg.policy, model=model)
Expand Down
15 changes: 10 additions & 5 deletions dizoo/box2d/lunarlander/entry/lunarlander_hpt_example.py
Original file line number Diff line number Diff line change
Expand Up @@ -21,16 +21,19 @@

def main():
logging.getLogger().setLevel(logging.INFO)
cfg = compile_config(main_config, create_cfg=create_config, auto=True, save_cfg=task.router.node_id == 0)
cfg = compile_config(main_config, create_cfg=create_config,
auto=True, save_cfg=task.router.node_id == 0)
ding_init(cfg)

with task.start(async_mode=False, ctx=OnlineRLContext()):
collector_env = SubprocessEnvManagerV2(
env_fn=[lambda: DingEnvWrapper(gym.make("LunarLander-v2")) for _ in range(cfg.env.collector_env_num)],
env_fn=[lambda: DingEnvWrapper(
gym.make("LunarLander-v2")) for _ in range(cfg.env.collector_env_num)],
cfg=cfg.env.manager
)
evaluator_env = SubprocessEnvManagerV2(
env_fn=[lambda: DingEnvWrapper(gym.make("LunarLander-v2")) for _ in range(cfg.env.evaluator_env_num)],
env_fn=[lambda: DingEnvWrapper(
gym.make("LunarLander-v2")) for _ in range(cfg.env.evaluator_env_num)],
cfg=cfg.env.manager
)

Expand All @@ -39,8 +42,10 @@ def main():
# Migrating models to the GPU
device = torch.device("cuda" if torch.cuda.is_available() else "cpu")
# HPT introduces a Policy Stem module, which processes the input features using Cross-Attention.
model = HPT(cfg.policy.model.obs_shape, cfg.policy.model.action_shape).to(device)
buffer_ = DequeBuffer(size=cfg.policy.other.replay_buffer.replay_buffer_size)
model = HPT(cfg.policy.model.obs_shape,
cfg.policy.model.action_shape).to(device)
buffer_ = DequeBuffer(
size=cfg.policy.other.replay_buffer.replay_buffer_size)

# Pass the model into Policy
policy = DQNPolicy(cfg.policy, model=model)
Expand Down

0 comments on commit 95a3edc

Please sign in to comment.